Skip to content

Commit a69863a

Browse files
authored
Merge pull request #2156 from huggingface/hiera
WIP Hiera implementation.
2 parents f7aa0a1 + 7a4e987 commit a69863a

40 files changed

+1002
-68
lines changed

tests/test_models.py

Lines changed: 4 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -52,15 +52,15 @@
5252
'vision_transformer', 'vision_transformer_sam', 'vision_transformer_hybrid', 'vision_transformer_relpos',
5353
'beit', 'mvitv2', 'eva', 'cait', 'xcit', 'volo', 'twins', 'deit', 'swin_transformer', 'swin_transformer_v2',
5454
'swin_transformer_v2_cr', 'maxxvit', 'efficientnet', 'mobilenetv3', 'levit', 'efficientformer', 'resnet',
55-
'regnet', 'byobnet', 'byoanet', 'mlp_mixer'
55+
'regnet', 'byobnet', 'byoanet', 'mlp_mixer', 'hiera',
5656
]
5757

5858
# transformer / hybrid models don't support full set of spatial / feature APIs and/or have spatial output.
5959
NON_STD_FILTERS = [
6060
'vit_*', 'tnt_*', 'pit_*', 'coat_*', 'cait_*', '*mixer_*', 'gmlp_*', 'resmlp_*', 'twins_*',
6161
'convit_*', 'levit*', 'visformer*', 'deit*', 'xcit_*', 'crossvit_*', 'beit*',
6262
'poolformer_*', 'volo_*', 'sequencer2d_*', 'mvitv2*', 'gcvit*', 'efficientformer*',
63-
'eva_*', 'flexivit*', 'eva02*', 'samvit_*', 'efficientvit_m*', 'tiny_vit_*'
63+
'eva_*', 'flexivit*', 'eva02*', 'samvit_*', 'efficientvit_m*', 'tiny_vit_*', 'hiera_*'
6464
]
6565
NUM_NON_STD = len(NON_STD_FILTERS)
6666

@@ -77,7 +77,7 @@
7777
EXCLUDE_FILTERS = ['*enormous*']
7878
NON_STD_EXCLUDE_FILTERS = ['*gigantic*', '*enormous*']
7979

80-
EXCLUDE_JIT_FILTERS = []
80+
EXCLUDE_JIT_FILTERS = ['hiera_*']
8181

8282
TARGET_FWD_SIZE = MAX_FWD_SIZE = 384
8383
TARGET_BWD_SIZE = 128
@@ -486,7 +486,7 @@ def _create_fx_model(model, train=False):
486486
return fx_model
487487

488488

489-
EXCLUDE_FX_FILTERS = ['vit_gi*']
489+
EXCLUDE_FX_FILTERS = ['vit_gi*', 'hiera*']
490490
# not enough memory to run fx on more models than other tests
491491
if 'GITHUB_ACTIONS' in os.environ:
492492
EXCLUDE_FX_FILTERS += [

timm/layers/classifier.py

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -108,7 +108,7 @@ def __init__(
108108
self.fc = fc
109109
self.flatten = nn.Flatten(1) if use_conv and pool_type else nn.Identity()
110110

111-
def reset(self, num_classes, pool_type=None):
111+
def reset(self, num_classes: int, pool_type: Optional[str] = None):
112112
if pool_type is not None and pool_type != self.global_pool.pool_type:
113113
self.global_pool, self.fc = create_classifier(
114114
self.in_features,
@@ -180,7 +180,7 @@ def __init__(
180180
self.drop = nn.Dropout(drop_rate)
181181
self.fc = linear_layer(self.num_features, num_classes) if num_classes > 0 else nn.Identity()
182182

183-
def reset(self, num_classes, pool_type=None):
183+
def reset(self, num_classes: int, pool_type: Optional[str] = None):
184184
if pool_type is not None:
185185
self.global_pool = SelectAdaptivePool2d(pool_type=pool_type)
186186
self.flatten = nn.Flatten(1) if pool_type else nn.Identity()

timm/layers/create_norm.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -47,7 +47,7 @@ def get_norm_layer(norm_layer):
4747
if isinstance(norm_layer, str):
4848
if not norm_layer:
4949
return None
50-
layer_name = norm_layer.replace('_', '')
50+
layer_name = norm_layer.replace('_', '').lower()
5151
norm_layer = _NORM_MAP[layer_name]
5252
else:
5353
norm_layer = norm_layer

timm/models/__init__.py

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -26,6 +26,7 @@
2626
from .ghostnet import *
2727
from .hardcorenas import *
2828
from .hgnet import *
29+
from .hiera import *
2930
from .hrnet import *
3031
from .inception_next import *
3132
from .inception_resnet_v2 import *

timm/models/beit.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -395,7 +395,7 @@ def group_matcher(self, coarse=False):
395395
def get_classifier(self):
396396
return self.head
397397

398-
def reset_classifier(self, num_classes, global_pool=None):
398+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
399399
self.num_classes = num_classes
400400
if global_pool is not None:
401401
self.global_pool = global_pool

timm/models/cait.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -331,7 +331,7 @@ def _matcher(name):
331331
def get_classifier(self):
332332
return self.head
333333

334-
def reset_classifier(self, num_classes, global_pool=None):
334+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
335335
self.num_classes = num_classes
336336
if global_pool is not None:
337337
assert global_pool in ('', 'token', 'avg')

timm/models/coat.py

Lines changed: 2 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -7,8 +7,7 @@
77
88
Modified from timm/models/vision_transformer.py
99
"""
10-
from functools import partial
11-
from typing import Tuple, List, Union
10+
from typing import List, Optional, Union, Tuple
1211

1312
import torch
1413
import torch.nn as nn
@@ -560,7 +559,7 @@ def group_matcher(self, coarse=False):
560559
def get_classifier(self):
561560
return self.head
562561

563-
def reset_classifier(self, num_classes, global_pool=None):
562+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
564563
self.num_classes = num_classes
565564
if global_pool is not None:
566565
assert global_pool in ('token', 'avg')

timm/models/convit.py

Lines changed: 2 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -21,8 +21,7 @@
2121
'''These modules are adapted from those of timm, see
2222
https://github.com/rwightman/pytorch-image-models/blob/master/timm/models/vision_transformer.py
2323
'''
24-
25-
from functools import partial
24+
from typing import Optional
2625

2726
import torch
2827
import torch.nn as nn
@@ -349,7 +348,7 @@ def set_grad_checkpointing(self, enable=True):
349348
def get_classifier(self):
350349
return self.head
351350

352-
def reset_classifier(self, num_classes, global_pool=None):
351+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
353352
self.num_classes = num_classes
354353
if global_pool is not None:
355354
assert global_pool in ('', 'token', 'avg')

timm/models/convmixer.py

Lines changed: 3 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -1,6 +1,8 @@
11
""" ConvMixer
22
33
"""
4+
from typing import Optional
5+
46
import torch
57
import torch.nn as nn
68

@@ -75,7 +77,7 @@ def set_grad_checkpointing(self, enable=True):
7577
def get_classifier(self):
7678
return self.head
7779

78-
def reset_classifier(self, num_classes, global_pool=None):
80+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
7981
self.num_classes = num_classes
8082
if global_pool is not None:
8183
self.pooling = SelectAdaptivePool2d(pool_type=global_pool, flatten=True)

timm/models/convnext.py

Lines changed: 0 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -37,7 +37,6 @@
3737
# LICENSE file in the root directory of this source tree (Attribution-NonCommercial 4.0 International (CC BY-NC 4.0))
3838
# No code was used directly from ConvNeXt-V2, however the weights are CC BY-NC 4.0 so beware if using commercially.
3939

40-
from collections import OrderedDict
4140
from functools import partial
4241
from typing import Callable, List, Optional, Tuple, Union
4342

timm/models/crossvit.py

Lines changed: 2 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -25,8 +25,7 @@
2525
2626
"""
2727
from functools import partial
28-
from typing import List
29-
from typing import Tuple
28+
from typing import List, Optional, Tuple
3029

3130
import torch
3231
import torch.hub
@@ -419,7 +418,7 @@ def set_grad_checkpointing(self, enable=True):
419418
def get_classifier(self):
420419
return self.head
421420

422-
def reset_classifier(self, num_classes, global_pool=None):
421+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
423422
self.num_classes = num_classes
424423
if global_pool is not None:
425424
assert global_pool in ('token', 'avg')

timm/models/davit.py

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -12,7 +12,7 @@
1212
# All rights reserved.
1313
# This source code is licensed under the MIT license
1414
from functools import partial
15-
from typing import Tuple
15+
from typing import Optional, Tuple
1616

1717
import torch
1818
import torch.nn as nn
@@ -568,7 +568,7 @@ def set_grad_checkpointing(self, enable=True):
568568
def get_classifier(self):
569569
return self.head.fc
570570

571-
def reset_classifier(self, num_classes, global_pool=None):
571+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
572572
self.head.reset(num_classes, global_pool)
573573

574574
def forward_features(self, x):

timm/models/deit.py

Lines changed: 2 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -11,7 +11,7 @@
1111
# Copyright (c) 2015-present, Facebook, Inc.
1212
# All rights reserved.
1313
from functools import partial
14-
from typing import Sequence, Union
14+
from typing import Optional
1515

1616
import torch
1717
from torch import nn as nn
@@ -20,7 +20,6 @@
2020
from timm.layers import resample_abs_pos_embed
2121
from timm.models.vision_transformer import VisionTransformer, trunc_normal_, checkpoint_filter_fn
2222
from ._builder import build_model_with_cfg
23-
from ._manipulate import checkpoint_seq
2423
from ._registry import generate_default_cfgs, register_model, register_model_deprecations
2524

2625
__all__ = ['VisionTransformerDistilled'] # model_registry will add each entrypoint fn to this
@@ -64,7 +63,7 @@ def group_matcher(self, coarse=False):
6463
def get_classifier(self):
6564
return self.head, self.head_dist
6665

67-
def reset_classifier(self, num_classes, global_pool=None):
66+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
6867
self.num_classes = num_classes
6968
self.head = nn.Linear(self.embed_dim, num_classes) if num_classes > 0 else nn.Identity()
7069
self.head_dist = nn.Linear(self.embed_dim, self.num_classes) if num_classes > 0 else nn.Identity()

timm/models/edgenext.py

Lines changed: 1 addition & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -8,7 +8,6 @@
88
Modifications and additions for timm by / Copyright 2022, Ross Wightman
99
"""
1010
import math
11-
from collections import OrderedDict
1211
from functools import partial
1312
from typing import Tuple
1413

@@ -17,7 +16,7 @@
1716
from torch import nn
1817

1918
from timm.data import IMAGENET_DEFAULT_MEAN, IMAGENET_DEFAULT_STD
20-
from timm.layers import trunc_normal_tf_, DropPath, LayerNorm2d, Mlp, SelectAdaptivePool2d, create_conv2d, \
19+
from timm.layers import trunc_normal_tf_, DropPath, LayerNorm2d, Mlp, create_conv2d, \
2120
use_fused_attn, NormMlpClassifierHead, ClassifierHead
2221
from ._builder import build_model_with_cfg
2322
from ._features_fx import register_notrace_module

timm/models/efficientformer.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -449,7 +449,7 @@ def set_grad_checkpointing(self, enable=True):
449449
def get_classifier(self):
450450
return self.head, self.head_dist
451451

452-
def reset_classifier(self, num_classes, global_pool=None):
452+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
453453
self.num_classes = num_classes
454454
if global_pool is not None:
455455
self.global_pool = global_pool

timm/models/efficientformer_v2.py

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -16,7 +16,7 @@
1616
"""
1717
import math
1818
from functools import partial
19-
from typing import Dict
19+
from typing import Dict, Optional
2020

2121
import torch
2222
import torch.nn as nn
@@ -612,7 +612,7 @@ def set_grad_checkpointing(self, enable=True):
612612
def get_classifier(self):
613613
return self.head, self.head_dist
614614

615-
def reset_classifier(self, num_classes, global_pool=None):
615+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
616616
self.num_classes = num_classes
617617
if global_pool is not None:
618618
self.global_pool = global_pool

timm/models/efficientvit_mit.py

Lines changed: 2 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -13,7 +13,6 @@
1313
import torch
1414
import torch.nn as nn
1515
import torch.nn.functional as F
16-
from torch.nn.modules.batchnorm import _BatchNorm
1716

1817
from timm.data import IMAGENET_DEFAULT_MEAN, IMAGENET_DEFAULT_STD
1918
from timm.layers import SelectAdaptivePool2d, create_conv2d, GELUTanh
@@ -740,7 +739,7 @@ def set_grad_checkpointing(self, enable=True):
740739
def get_classifier(self):
741740
return self.head.classifier[-1]
742741

743-
def reset_classifier(self, num_classes, global_pool=None):
742+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
744743
self.num_classes = num_classes
745744
if global_pool is not None:
746745
self.global_pool = global_pool
@@ -858,7 +857,7 @@ def set_grad_checkpointing(self, enable=True):
858857
def get_classifier(self):
859858
return self.head.classifier[-1]
860859

861-
def reset_classifier(self, num_classes, global_pool=None):
860+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
862861
self.num_classes = num_classes
863862
if global_pool is not None:
864863
self.global_pool = global_pool

timm/models/efficientvit_msra.py

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -9,7 +9,7 @@
99
__all__ = ['EfficientVitMsra']
1010
import itertools
1111
from collections import OrderedDict
12-
from typing import Dict
12+
from typing import Dict, Optional
1313

1414
import torch
1515
import torch.nn as nn
@@ -464,7 +464,7 @@ def set_grad_checkpointing(self, enable=True):
464464
def get_classifier(self):
465465
return self.head.linear
466466

467-
def reset_classifier(self, num_classes, global_pool=None):
467+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
468468
self.num_classes = num_classes
469469
if global_pool is not None:
470470
if global_pool == 'avg':

timm/models/eva.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -539,7 +539,7 @@ def group_matcher(self, coarse=False):
539539
def get_classifier(self):
540540
return self.head
541541

542-
def reset_classifier(self, num_classes, global_pool=None):
542+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
543543
self.num_classes = num_classes
544544
if global_pool is not None:
545545
self.global_pool = global_pool

timm/models/fastvit.py

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -396,7 +396,7 @@ def reparameterize(self) -> None:
396396

397397
@staticmethod
398398
def _fuse_bn(
399-
conv: torch.Tensor, bn: nn.BatchNorm2d
399+
conv: nn.Conv2d, bn: nn.BatchNorm2d
400400
) -> Tuple[torch.Tensor, torch.Tensor]:
401401
"""Method to fuse batchnorm layer with conv layer.
402402
@@ -1232,7 +1232,7 @@ def set_grad_checkpointing(self, enable=True):
12321232
def get_classifier(self):
12331233
return self.head.fc
12341234

1235-
def reset_classifier(self, num_classes, global_pool=None):
1235+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
12361236
self.num_classes = num_classes
12371237
self.head.reset(num_classes, global_pool)
12381238

timm/models/focalnet.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -454,7 +454,7 @@ def set_grad_checkpointing(self, enable=True):
454454
def get_classifier(self):
455455
return self.head.fc
456456

457-
def reset_classifier(self, num_classes, global_pool=None):
457+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
458458
self.head.reset(num_classes, pool_type=global_pool)
459459

460460
def forward_features(self, x):

timm/models/gcvit.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -489,7 +489,7 @@ def set_grad_checkpointing(self, enable=True):
489489
def get_classifier(self):
490490
return self.head.fc
491491

492-
def reset_classifier(self, num_classes, global_pool=None):
492+
def reset_classifier(self, num_classes: int, global_pool: Optional[str] = None):
493493
self.num_classes = num_classes
494494
if global_pool is None:
495495
global_pool = self.head.global_pool.pool_type

0 commit comments

Comments
 (0)