modeling_resnet.py 19 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517
  1. # coding=utf-8
  2. # Copyright 2022 Microsoft Research, Inc. and The HuggingFace Inc. team. All rights reserved.
  3. #
  4. # Licensed under the Apache License, Version 2.0 (the "License");
  5. # you may not use this file except in compliance with the License.
  6. # You may obtain a copy of the License at
  7. #
  8. # http://www.apache.org/licenses/LICENSE-2.0
  9. #
  10. # Unless required by applicable law or agreed to in writing, software
  11. # distributed under the License is distributed on an "AS IS" BASIS,
  12. # WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
  13. # See the License for the specific language governing permissions and
  14. # limitations under the License.
  15. """PyTorch ResNet model."""
  16. import math
  17. from typing import Optional
  18. import torch
  19. import torch.utils.checkpoint
  20. from torch import Tensor, nn
  21. from torch.nn import BCEWithLogitsLoss, CrossEntropyLoss, MSELoss
  22. from ...activations import ACT2FN
  23. from ...modeling_outputs import (
  24. BackboneOutput,
  25. BaseModelOutputWithNoAttention,
  26. BaseModelOutputWithPoolingAndNoAttention,
  27. ImageClassifierOutputWithNoAttention,
  28. )
  29. from ...modeling_utils import PreTrainedModel
  30. from ...utils import (
  31. add_code_sample_docstrings,
  32. add_start_docstrings,
  33. add_start_docstrings_to_model_forward,
  34. logging,
  35. replace_return_docstrings,
  36. )
  37. from ...utils.backbone_utils import BackboneMixin
  38. from .configuration_resnet import ResNetConfig
  39. logger = logging.get_logger(__name__)
  40. # General docstring
  41. _CONFIG_FOR_DOC = "ResNetConfig"
  42. # Base docstring
  43. _CHECKPOINT_FOR_DOC = "microsoft/resnet-50"
  44. _EXPECTED_OUTPUT_SHAPE = [1, 2048, 7, 7]
  45. # Image classification docstring
  46. _IMAGE_CLASS_CHECKPOINT = "microsoft/resnet-50"
  47. _IMAGE_CLASS_EXPECTED_OUTPUT = "tiger cat"
  48. class ResNetConvLayer(nn.Module):
  49. def __init__(
  50. self, in_channels: int, out_channels: int, kernel_size: int = 3, stride: int = 1, activation: str = "relu"
  51. ):
  52. super().__init__()
  53. self.convolution = nn.Conv2d(
  54. in_channels, out_channels, kernel_size=kernel_size, stride=stride, padding=kernel_size // 2, bias=False
  55. )
  56. self.normalization = nn.BatchNorm2d(out_channels)
  57. self.activation = ACT2FN[activation] if activation is not None else nn.Identity()
  58. def forward(self, input: Tensor) -> Tensor:
  59. hidden_state = self.convolution(input)
  60. hidden_state = self.normalization(hidden_state)
  61. hidden_state = self.activation(hidden_state)
  62. return hidden_state
  63. class ResNetEmbeddings(nn.Module):
  64. """
  65. ResNet Embeddings (stem) composed of a single aggressive convolution.
  66. """
  67. def __init__(self, config: ResNetConfig):
  68. super().__init__()
  69. self.embedder = ResNetConvLayer(
  70. config.num_channels, config.embedding_size, kernel_size=7, stride=2, activation=config.hidden_act
  71. )
  72. self.pooler = nn.MaxPool2d(kernel_size=3, stride=2, padding=1)
  73. self.num_channels = config.num_channels
  74. def forward(self, pixel_values: Tensor) -> Tensor:
  75. num_channels = pixel_values.shape[1]
  76. if num_channels != self.num_channels:
  77. raise ValueError(
  78. "Make sure that the channel dimension of the pixel values match with the one set in the configuration."
  79. )
  80. embedding = self.embedder(pixel_values)
  81. embedding = self.pooler(embedding)
  82. return embedding
  83. class ResNetShortCut(nn.Module):
  84. """
  85. ResNet shortcut, used to project the residual features to the correct size. If needed, it is also used to
  86. downsample the input using `stride=2`.
  87. """
  88. def __init__(self, in_channels: int, out_channels: int, stride: int = 2):
  89. super().__init__()
  90. self.convolution = nn.Conv2d(in_channels, out_channels, kernel_size=1, stride=stride, bias=False)
  91. self.normalization = nn.BatchNorm2d(out_channels)
  92. def forward(self, input: Tensor) -> Tensor:
  93. hidden_state = self.convolution(input)
  94. hidden_state = self.normalization(hidden_state)
  95. return hidden_state
  96. class ResNetBasicLayer(nn.Module):
  97. """
  98. A classic ResNet's residual layer composed by two `3x3` convolutions.
  99. """
  100. def __init__(self, in_channels: int, out_channels: int, stride: int = 1, activation: str = "relu"):
  101. super().__init__()
  102. should_apply_shortcut = in_channels != out_channels or stride != 1
  103. self.shortcut = (
  104. ResNetShortCut(in_channels, out_channels, stride=stride) if should_apply_shortcut else nn.Identity()
  105. )
  106. self.layer = nn.Sequential(
  107. ResNetConvLayer(in_channels, out_channels, stride=stride),
  108. ResNetConvLayer(out_channels, out_channels, activation=None),
  109. )
  110. self.activation = ACT2FN[activation]
  111. def forward(self, hidden_state):
  112. residual = hidden_state
  113. hidden_state = self.layer(hidden_state)
  114. residual = self.shortcut(residual)
  115. hidden_state += residual
  116. hidden_state = self.activation(hidden_state)
  117. return hidden_state
  118. class ResNetBottleNeckLayer(nn.Module):
  119. """
  120. A classic ResNet's bottleneck layer composed by three `3x3` convolutions.
  121. The first `1x1` convolution reduces the input by a factor of `reduction` in order to make the second `3x3`
  122. convolution faster. The last `1x1` convolution remaps the reduced features to `out_channels`. If
  123. `downsample_in_bottleneck` is true, downsample will be in the first layer instead of the second layer.
  124. """
  125. def __init__(
  126. self,
  127. in_channels: int,
  128. out_channels: int,
  129. stride: int = 1,
  130. activation: str = "relu",
  131. reduction: int = 4,
  132. downsample_in_bottleneck: bool = False,
  133. ):
  134. super().__init__()
  135. should_apply_shortcut = in_channels != out_channels or stride != 1
  136. reduces_channels = out_channels // reduction
  137. self.shortcut = (
  138. ResNetShortCut(in_channels, out_channels, stride=stride) if should_apply_shortcut else nn.Identity()
  139. )
  140. self.layer = nn.Sequential(
  141. ResNetConvLayer(
  142. in_channels, reduces_channels, kernel_size=1, stride=stride if downsample_in_bottleneck else 1
  143. ),
  144. ResNetConvLayer(reduces_channels, reduces_channels, stride=stride if not downsample_in_bottleneck else 1),
  145. ResNetConvLayer(reduces_channels, out_channels, kernel_size=1, activation=None),
  146. )
  147. self.activation = ACT2FN[activation]
  148. def forward(self, hidden_state):
  149. residual = hidden_state
  150. hidden_state = self.layer(hidden_state)
  151. residual = self.shortcut(residual)
  152. hidden_state += residual
  153. hidden_state = self.activation(hidden_state)
  154. return hidden_state
  155. class ResNetStage(nn.Module):
  156. """
  157. A ResNet stage composed by stacked layers.
  158. """
  159. def __init__(
  160. self,
  161. config: ResNetConfig,
  162. in_channels: int,
  163. out_channels: int,
  164. stride: int = 2,
  165. depth: int = 2,
  166. ):
  167. super().__init__()
  168. layer = ResNetBottleNeckLayer if config.layer_type == "bottleneck" else ResNetBasicLayer
  169. if config.layer_type == "bottleneck":
  170. first_layer = layer(
  171. in_channels,
  172. out_channels,
  173. stride=stride,
  174. activation=config.hidden_act,
  175. downsample_in_bottleneck=config.downsample_in_bottleneck,
  176. )
  177. else:
  178. first_layer = layer(in_channels, out_channels, stride=stride, activation=config.hidden_act)
  179. self.layers = nn.Sequential(
  180. first_layer, *[layer(out_channels, out_channels, activation=config.hidden_act) for _ in range(depth - 1)]
  181. )
  182. def forward(self, input: Tensor) -> Tensor:
  183. hidden_state = input
  184. for layer in self.layers:
  185. hidden_state = layer(hidden_state)
  186. return hidden_state
  187. class ResNetEncoder(nn.Module):
  188. def __init__(self, config: ResNetConfig):
  189. super().__init__()
  190. self.stages = nn.ModuleList([])
  191. # based on `downsample_in_first_stage` the first layer of the first stage may or may not downsample the input
  192. self.stages.append(
  193. ResNetStage(
  194. config,
  195. config.embedding_size,
  196. config.hidden_sizes[0],
  197. stride=2 if config.downsample_in_first_stage else 1,
  198. depth=config.depths[0],
  199. )
  200. )
  201. in_out_channels = zip(config.hidden_sizes, config.hidden_sizes[1:])
  202. for (in_channels, out_channels), depth in zip(in_out_channels, config.depths[1:]):
  203. self.stages.append(ResNetStage(config, in_channels, out_channels, depth=depth))
  204. def forward(
  205. self, hidden_state: Tensor, output_hidden_states: bool = False, return_dict: bool = True
  206. ) -> BaseModelOutputWithNoAttention:
  207. hidden_states = () if output_hidden_states else None
  208. for stage_module in self.stages:
  209. if output_hidden_states:
  210. hidden_states = hidden_states + (hidden_state,)
  211. hidden_state = stage_module(hidden_state)
  212. if output_hidden_states:
  213. hidden_states = hidden_states + (hidden_state,)
  214. if not return_dict:
  215. return tuple(v for v in [hidden_state, hidden_states] if v is not None)
  216. return BaseModelOutputWithNoAttention(
  217. last_hidden_state=hidden_state,
  218. hidden_states=hidden_states,
  219. )
  220. class ResNetPreTrainedModel(PreTrainedModel):
  221. """
  222. An abstract class to handle weights initialization and a simple interface for downloading and loading pretrained
  223. models.
  224. """
  225. config_class = ResNetConfig
  226. base_model_prefix = "resnet"
  227. main_input_name = "pixel_values"
  228. _no_split_modules = ["ResNetConvLayer", "ResNetShortCut"]
  229. def _init_weights(self, module):
  230. if isinstance(module, nn.Conv2d):
  231. nn.init.kaiming_normal_(module.weight, mode="fan_out", nonlinearity="relu")
  232. # copied from the `reset_parameters` method of `class Linear(Module)` in `torch`.
  233. elif isinstance(module, nn.Linear):
  234. nn.init.kaiming_uniform_(module.weight, a=math.sqrt(5))
  235. if module.bias is not None:
  236. fan_in, _ = nn.init._calculate_fan_in_and_fan_out(module.weight)
  237. bound = 1 / math.sqrt(fan_in) if fan_in > 0 else 0
  238. nn.init.uniform_(module.bias, -bound, bound)
  239. elif isinstance(module, (nn.BatchNorm2d, nn.GroupNorm)):
  240. nn.init.constant_(module.weight, 1)
  241. nn.init.constant_(module.bias, 0)
  242. RESNET_START_DOCSTRING = r"""
  243. This model is a PyTorch [torch.nn.Module](https://pytorch.org/docs/stable/nn.html#torch.nn.Module) subclass. Use it
  244. as a regular PyTorch Module and refer to the PyTorch documentation for all matter related to general usage and
  245. behavior.
  246. Parameters:
  247. config ([`ResNetConfig`]): Model configuration class with all the parameters of the model.
  248. Initializing with a config file does not load the weights associated with the model, only the
  249. configuration. Check out the [`~PreTrainedModel.from_pretrained`] method to load the model weights.
  250. """
  251. RESNET_INPUTS_DOCSTRING = r"""
  252. Args:
  253. pixel_values (`torch.FloatTensor` of shape `(batch_size, num_channels, height, width)`):
  254. Pixel values. Pixel values can be obtained using [`AutoImageProcessor`]. See
  255. [`ConvNextImageProcessor.__call__`] for details.
  256. output_hidden_states (`bool`, *optional*):
  257. Whether or not to return the hidden states of all layers. See `hidden_states` under returned tensors for
  258. more detail.
  259. return_dict (`bool`, *optional*):
  260. Whether or not to return a [`~utils.ModelOutput`] instead of a plain tuple.
  261. """
  262. @add_start_docstrings(
  263. "The bare ResNet model outputting raw features without any specific head on top.",
  264. RESNET_START_DOCSTRING,
  265. )
  266. class ResNetModel(ResNetPreTrainedModel):
  267. def __init__(self, config):
  268. super().__init__(config)
  269. self.config = config
  270. self.embedder = ResNetEmbeddings(config)
  271. self.encoder = ResNetEncoder(config)
  272. self.pooler = nn.AdaptiveAvgPool2d((1, 1))
  273. # Initialize weights and apply final processing
  274. self.post_init()
  275. @add_start_docstrings_to_model_forward(RESNET_INPUTS_DOCSTRING)
  276. @add_code_sample_docstrings(
  277. checkpoint=_CHECKPOINT_FOR_DOC,
  278. output_type=BaseModelOutputWithPoolingAndNoAttention,
  279. config_class=_CONFIG_FOR_DOC,
  280. modality="vision",
  281. expected_output=_EXPECTED_OUTPUT_SHAPE,
  282. )
  283. def forward(
  284. self, pixel_values: Tensor, output_hidden_states: Optional[bool] = None, return_dict: Optional[bool] = None
  285. ) -> BaseModelOutputWithPoolingAndNoAttention:
  286. output_hidden_states = (
  287. output_hidden_states if output_hidden_states is not None else self.config.output_hidden_states
  288. )
  289. return_dict = return_dict if return_dict is not None else self.config.use_return_dict
  290. embedding_output = self.embedder(pixel_values)
  291. encoder_outputs = self.encoder(
  292. embedding_output, output_hidden_states=output_hidden_states, return_dict=return_dict
  293. )
  294. last_hidden_state = encoder_outputs[0]
  295. pooled_output = self.pooler(last_hidden_state)
  296. if not return_dict:
  297. return (last_hidden_state, pooled_output) + encoder_outputs[1:]
  298. return BaseModelOutputWithPoolingAndNoAttention(
  299. last_hidden_state=last_hidden_state,
  300. pooler_output=pooled_output,
  301. hidden_states=encoder_outputs.hidden_states,
  302. )
  303. @add_start_docstrings(
  304. """
  305. ResNet Model with an image classification head on top (a linear layer on top of the pooled features), e.g. for
  306. ImageNet.
  307. """,
  308. RESNET_START_DOCSTRING,
  309. )
  310. class ResNetForImageClassification(ResNetPreTrainedModel):
  311. def __init__(self, config):
  312. super().__init__(config)
  313. self.num_labels = config.num_labels
  314. self.resnet = ResNetModel(config)
  315. # classification head
  316. self.classifier = nn.Sequential(
  317. nn.Flatten(),
  318. nn.Linear(config.hidden_sizes[-1], config.num_labels) if config.num_labels > 0 else nn.Identity(),
  319. )
  320. # initialize weights and apply final processing
  321. self.post_init()
  322. @add_start_docstrings_to_model_forward(RESNET_INPUTS_DOCSTRING)
  323. @add_code_sample_docstrings(
  324. checkpoint=_IMAGE_CLASS_CHECKPOINT,
  325. output_type=ImageClassifierOutputWithNoAttention,
  326. config_class=_CONFIG_FOR_DOC,
  327. expected_output=_IMAGE_CLASS_EXPECTED_OUTPUT,
  328. )
  329. def forward(
  330. self,
  331. pixel_values: Optional[torch.FloatTensor] = None,
  332. labels: Optional[torch.LongTensor] = None,
  333. output_hidden_states: Optional[bool] = None,
  334. return_dict: Optional[bool] = None,
  335. ) -> ImageClassifierOutputWithNoAttention:
  336. r"""
  337. labels (`torch.LongTensor` of shape `(batch_size,)`, *optional*):
  338. Labels for computing the image classification/regression loss. Indices should be in `[0, ...,
  339. config.num_labels - 1]`. If `config.num_labels > 1` a classification loss is computed (Cross-Entropy).
  340. """
  341. return_dict = return_dict if return_dict is not None else self.config.use_return_dict
  342. outputs = self.resnet(pixel_values, output_hidden_states=output_hidden_states, return_dict=return_dict)
  343. pooled_output = outputs.pooler_output if return_dict else outputs[1]
  344. logits = self.classifier(pooled_output)
  345. loss = None
  346. if labels is not None:
  347. if self.config.problem_type is None:
  348. if self.num_labels == 1:
  349. self.config.problem_type = "regression"
  350. elif self.num_labels > 1 and (labels.dtype == torch.long or labels.dtype == torch.int):
  351. self.config.problem_type = "single_label_classification"
  352. else:
  353. self.config.problem_type = "multi_label_classification"
  354. if self.config.problem_type == "regression":
  355. loss_fct = MSELoss()
  356. if self.num_labels == 1:
  357. loss = loss_fct(logits.squeeze(), labels.squeeze())
  358. else:
  359. loss = loss_fct(logits, labels)
  360. elif self.config.problem_type == "single_label_classification":
  361. loss_fct = CrossEntropyLoss()
  362. loss = loss_fct(logits.view(-1, self.num_labels), labels.view(-1))
  363. elif self.config.problem_type == "multi_label_classification":
  364. loss_fct = BCEWithLogitsLoss()
  365. loss = loss_fct(logits, labels)
  366. if not return_dict:
  367. output = (logits,) + outputs[2:]
  368. return (loss,) + output if loss is not None else output
  369. return ImageClassifierOutputWithNoAttention(loss=loss, logits=logits, hidden_states=outputs.hidden_states)
  370. @add_start_docstrings(
  371. """
  372. ResNet backbone, to be used with frameworks like DETR and MaskFormer.
  373. """,
  374. RESNET_START_DOCSTRING,
  375. )
  376. class ResNetBackbone(ResNetPreTrainedModel, BackboneMixin):
  377. def __init__(self, config):
  378. super().__init__(config)
  379. super()._init_backbone(config)
  380. self.num_features = [config.embedding_size] + config.hidden_sizes
  381. self.embedder = ResNetEmbeddings(config)
  382. self.encoder = ResNetEncoder(config)
  383. # initialize weights and apply final processing
  384. self.post_init()
  385. @add_start_docstrings_to_model_forward(RESNET_INPUTS_DOCSTRING)
  386. @replace_return_docstrings(output_type=BackboneOutput, config_class=_CONFIG_FOR_DOC)
  387. def forward(
  388. self, pixel_values: Tensor, output_hidden_states: Optional[bool] = None, return_dict: Optional[bool] = None
  389. ) -> BackboneOutput:
  390. """
  391. Returns:
  392. Examples:
  393. ```python
  394. >>> from transformers import AutoImageProcessor, AutoBackbone
  395. >>> import torch
  396. >>> from PIL import Image
  397. >>> import requests
  398. >>> url = "http://images.cocodataset.org/val2017/000000039769.jpg"
  399. >>> image = Image.open(requests.get(url, stream=True).raw)
  400. >>> processor = AutoImageProcessor.from_pretrained("microsoft/resnet-50")
  401. >>> model = AutoBackbone.from_pretrained(
  402. ... "microsoft/resnet-50", out_features=["stage1", "stage2", "stage3", "stage4"]
  403. ... )
  404. >>> inputs = processor(image, return_tensors="pt")
  405. >>> outputs = model(**inputs)
  406. >>> feature_maps = outputs.feature_maps
  407. >>> list(feature_maps[-1].shape)
  408. [1, 2048, 7, 7]
  409. ```"""
  410. return_dict = return_dict if return_dict is not None else self.config.use_return_dict
  411. output_hidden_states = (
  412. output_hidden_states if output_hidden_states is not None else self.config.output_hidden_states
  413. )
  414. embedding_output = self.embedder(pixel_values)
  415. outputs = self.encoder(embedding_output, output_hidden_states=True, return_dict=True)
  416. hidden_states = outputs.hidden_states
  417. feature_maps = ()
  418. for idx, stage in enumerate(self.stage_names):
  419. if stage in self.out_features:
  420. feature_maps += (hidden_states[idx],)
  421. if not return_dict:
  422. output = (feature_maps,)
  423. if output_hidden_states:
  424. output += (outputs.hidden_states,)
  425. return output
  426. return BackboneOutput(
  427. feature_maps=feature_maps,
  428. hidden_states=outputs.hidden_states if output_hidden_states else None,
  429. attentions=None,
  430. )