Skip to content

Commit dd7ac4c

Browse files
authored
[tests] Really use small models in all fast tests (#40945)
* start * xcodec * chameleon * start * layoutlm2 * layoutlm * remove skip * oups * timm_wrapper * add default * doc * consistency
1 parent 2ce35a2 commit dd7ac4c

File tree

9 files changed

+34
-45
lines changed

9 files changed

+34
-45
lines changed

src/transformers/models/gemma3n/configuration_gemma3n.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -502,10 +502,10 @@ def __init__(
502502
**kwargs,
503503
):
504504
super().__init__(**kwargs)
505+
self.architecture = architecture
505506
self.initializer_range = initializer_range
506507
self.do_pooling = do_pooling
507508
self.model_args = model_args # named "model_args" for BC with timm
508-
self.architecture = architecture
509509
self.hidden_size = hidden_size
510510
self.vocab_size = vocab_size
511511
self.vocab_offset = vocab_offset

src/transformers/models/timm_wrapper/configuration_timm_wrapper.py

Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -41,6 +41,8 @@ class TimmWrapperConfig(PretrainedConfig):
4141
imagenet models is set to `None` due to occlusions in the label descriptions.
4242
4343
Args:
44+
architecture (`str`, *optional*, defaults to `"resnet50"`):
45+
The timm architecture to load.
4446
initializer_range (`float`, *optional*, defaults to 0.02):
4547
The standard deviation of the truncated_normal_initializer for initializing all weight matrices.
4648
do_pooling (`bool`, *optional*, defaults to `True`):
@@ -65,11 +67,13 @@ class TimmWrapperConfig(PretrainedConfig):
6567

6668
def __init__(
6769
self,
70+
architecture: str = "resnet50",
6871
initializer_range: float = 0.02,
6972
do_pooling: bool = True,
7073
model_args: Optional[dict[str, Any]] = None,
7174
**kwargs,
7275
):
76+
self.architecture = architecture
7377
self.initializer_range = initializer_range
7478
self.do_pooling = do_pooling
7579
self.model_args = model_args # named "model_args" for BC with timm

tests/models/chameleon/test_modeling_chameleon.py

Lines changed: 1 addition & 9 deletions
Original file line numberDiff line numberDiff line change
@@ -76,7 +76,7 @@ def __init__(
7676
pad_token_id=0,
7777
vq_num_embeds=5,
7878
vq_embed_dim=5,
79-
vq_channel_multiplier=[1, 4],
79+
vq_channel_multiplier=[1, 2],
8080
vq_img_token_start_id=10, # has to be less than vocab size when added with vq_num_embeds
8181
scope=None,
8282
):
@@ -255,10 +255,6 @@ def test_model_rope_scaling(self, scaling_type):
255255
def test_batching_equivalence(self):
256256
pass
257257

258-
@unittest.skip("Chameleon VQ model cannot be squishes more due to hardcoded layer params in model code")
259-
def test_model_is_small(self):
260-
pass
261-
262258

263259
class ChameleonVision2SeqModelTester(ChameleonModelTester):
264260
def __init__(self, parent, image_size=10, **kwargs):
@@ -321,10 +317,6 @@ def test_disk_offload_bin(self):
321317
def test_disk_offload_safetensors(self):
322318
pass
323319

324-
@unittest.skip("Chameleon VQ model cannot be squishes more due to hardcoded layer params in model code")
325-
def test_model_is_small(self):
326-
pass
327-
328320
@unittest.skip("Chameleon applies key/query norm which doesn't work with packing")
329321
def test_flash_attention_2_padding_matches_padding_free_with_position_ids(self):
330322
pass

tests/models/emu3/test_modeling_emu3.py

Lines changed: 0 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -359,10 +359,6 @@ def test_initialization(self):
359359
def test_generate_with_static_cache(self):
360360
pass
361361

362-
# @unittest.skip("Emu3 can't be smaller than currently if we want to downsample images")
363-
# def test_model_is_small(self):
364-
# pass
365-
366362

367363
@require_torch
368364
class Emu3IntegrationTest(unittest.TestCase):

tests/models/layoutlmv2/test_modeling_layoutlmv2.py

Lines changed: 10 additions & 10 deletions
Original file line numberDiff line numberDiff line change
@@ -70,7 +70,7 @@ def __init__(
7070
type_vocab_size=16,
7171
type_sequence_label_size=2,
7272
initializer_range=0.02,
73-
image_feature_pool_shape=[7, 7, 256],
73+
image_feature_pool_shape=[7, 7, 32],
7474
coordinate_size=6,
7575
shape_size=6,
7676
num_labels=3,
@@ -106,6 +106,14 @@ def __init__(
106106
self.num_choices = num_choices
107107
self.scope = scope
108108
self.range_bbox = range_bbox
109+
detectron2_config = LayoutLMv2Config.get_default_detectron2_config()
110+
# We need to make the model smaller
111+
detectron2_config["MODEL.RESNETS.DEPTH"] = 50
112+
detectron2_config["MODEL.RESNETS.RES2_OUT_CHANNELS"] = 4
113+
detectron2_config["MODEL.RESNETS.STEM_OUT_CHANNELS"] = 4
114+
detectron2_config["MODEL.FPN.OUT_CHANNELS"] = 32
115+
detectron2_config["MODEL.RESNETS.NUM_GROUPS"] = 1
116+
self.detectron2_config = detectron2_config
109117

110118
def prepare_config_and_inputs(self):
111119
input_ids = ids_tensor([self.batch_size, self.seq_length], self.vocab_size)
@@ -158,13 +166,9 @@ def prepare_config_and_inputs(self):
158166
image_feature_pool_shape=self.image_feature_pool_shape,
159167
coordinate_size=self.coordinate_size,
160168
shape_size=self.shape_size,
169+
detectron2_config_args=self.detectron2_config,
161170
)
162171

163-
# use smaller resnet backbone to make tests faster
164-
config.detectron2_config_args["MODEL.RESNETS.DEPTH"] = 18
165-
config.detectron2_config_args["MODEL.RESNETS.RES2_OUT_CHANNELS"] = 64
166-
config.detectron2_config_args["MODEL.RESNETS.NUM_GROUPS"] = 1
167-
168172
return config, input_ids, bbox, image, token_type_ids, input_mask, sequence_labels, token_labels
169173

170174
def create_and_check_model(
@@ -422,10 +426,6 @@ def check_hidden_states_output(inputs_dict, config, model_class):
422426

423427
check_hidden_states_output(inputs_dict, config, model_class)
424428

425-
@unittest.skip(reason="We cannot configure detectron2 to output a smaller backbone")
426-
def test_model_is_small(self):
427-
pass
428-
429429
@slow
430430
def test_model_from_pretrained(self):
431431
model_name = "microsoft/layoutlmv2-base-uncased"

tests/models/qwen2_5_vl/test_modeling_qwen2_5_vl.py

Lines changed: 0 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -441,10 +441,6 @@ def test_sdpa_can_dispatch_on_flash(self):
441441
def test_multi_gpu_data_parallel_forward(self):
442442
pass
443443

444-
@unittest.skip(reason="We cannot configure to output a smaller model.")
445-
def test_model_is_small(self):
446-
pass
447-
448444

449445
@require_torch
450446
class Qwen2_5_VLIntegrationTest(unittest.TestCase):

tests/models/qwen2_vl/test_modeling_qwen2_vl.py

Lines changed: 0 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -394,10 +394,6 @@ def test_sdpa_can_dispatch_on_flash(self):
394394
def test_multi_gpu_data_parallel_forward(self):
395395
pass
396396

397-
@unittest.skip(reason="We cannot configure to output a smaller model.")
398-
def test_model_is_small(self):
399-
pass
400-
401397

402398
@require_torch
403399
class Qwen2VLIntegrationTest(unittest.TestCase):

tests/models/timm_wrapper/test_modeling_timm_wrapper.py

Lines changed: 4 additions & 7 deletions
Original file line numberDiff line numberDiff line change
@@ -53,14 +53,15 @@ class TimmWrapperModelTester:
5353
def __init__(
5454
self,
5555
parent,
56-
model_name="timm/resnet18.a1_in1k",
5756
batch_size=3,
5857
image_size=32,
5958
num_channels=3,
6059
is_training=True,
6160
):
6261
self.parent = parent
63-
self.model_name = model_name
62+
self.architecture = "resnet26"
63+
# We need this to make the model smaller
64+
self.model_args = {"channels": (16, 16, 16, 16)}
6465
self.batch_size = batch_size
6566
self.image_size = image_size
6667
self.num_channels = num_channels
@@ -73,7 +74,7 @@ def prepare_config_and_inputs(self):
7374
return config, pixel_values
7475

7576
def get_config(self):
76-
return TimmWrapperConfig.from_pretrained(self.model_name)
77+
return TimmWrapperConfig(architecture=self.architecture, model_args=self.model_args)
7778

7879
def prepare_config_and_inputs_for_common(self):
7980
config_and_inputs = self.prepare_config_and_inputs()
@@ -166,10 +167,6 @@ def test_initialization(self):
166167
def test_mismatched_shapes_have_properly_initialized_weights(self):
167168
pass
168169

169-
@unittest.skip(reason="Need to use a timm model and there is no tiny model available.")
170-
def test_model_is_small(self):
171-
pass
172-
173170
def test_gradient_checkpointing(self):
174171
config, _ = self.model_tester.prepare_config_and_inputs_for_common()
175172
model = TimmWrapperModel._from_config(config)

tests/models/xcodec/test_modeling_xcodec.py

Lines changed: 14 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -39,7 +39,7 @@
3939
if is_torch_available():
4040
import torch
4141

42-
from transformers import XcodecModel
42+
from transformers import DacConfig, HubertConfig, XcodecModel
4343

4444

4545
@require_torch
@@ -51,7 +51,7 @@ def __init__(
5151
num_channels=1,
5252
sample_rate=16000,
5353
codebook_size=1024,
54-
num_samples=400,
54+
num_samples=256,
5555
is_training=False,
5656
):
5757
self.parent = parent
@@ -61,6 +61,16 @@ def __init__(
6161
self.codebook_size = codebook_size
6262
self.is_training = is_training
6363
self.num_samples = num_samples
64+
self.acoustic_model_config = DacConfig(
65+
decoder_hidden_size=8, encoder_hidden_size=8, codebook_size=16, downsampling_ratios=[16, 16]
66+
)
67+
self.semantic_model_config = HubertConfig(
68+
hidden_size=32,
69+
num_hidden_layers=2,
70+
num_attention_heads=2,
71+
intermediate_size=12,
72+
conv_dim=(4, 4, 4, 4, 4, 4, 4),
73+
)
6474

6575
def prepare_config_and_inputs(self):
6676
config = self.get_config()
@@ -86,6 +96,8 @@ def get_config(self):
8696
sample_rate=self.sample_rate,
8797
audio_channels=self.num_channels,
8898
codebook_size=self.codebook_size,
99+
acoustic_model_config=self.acoustic_model_config,
100+
semantic_model_config=self.semantic_model_config,
89101
)
90102

91103
def create_and_check_model_forward(self, config, inputs_dict):
@@ -151,10 +163,6 @@ def test_gradient_checkpointing_backward_compatibility(self):
151163
model = model_class(config)
152164
self.assertTrue(model.is_gradient_checkpointing)
153165

154-
@unittest.skip(reason="We cannot configure to output a smaller model.")
155-
def test_model_is_small(self):
156-
pass
157-
158166
@unittest.skip(reason="The XcodecModel does not have `inputs_embeds` logics")
159167
def test_inputs_embeds(self):
160168
pass

0 commit comments

Comments
 (0)