You can not select more than 25 topics Topics must start with a chinese character,a letter or number, can include dashes ('-') and can be up to 35 characters long.

module.py 27 kB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728729730731732
  1. # MegEngine is Licensed under the Apache License, Version 2.0 (the "License")
  2. #
  3. # Copyright (c) 2014-2021 Megvii Inc. All rights reserved.
  4. #
  5. # Unless required by applicable law or agreed to in writing,
  6. # software distributed under the License is distributed on an
  7. # "AS IS" BASIS, WITHOUT ARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
  8. from abc import ABCMeta, abstractmethod
  9. from collections import OrderedDict
  10. from typing import Any, Callable, Iterable, Optional, Set, Tuple, Union
  11. import numpy as np
  12. from ..core.tensor.utils import make_shape_tuple
  13. from ..logger import get_logger
  14. from ..tensor import Parameter, Tensor
  15. from ..utils.deprecation import deprecated
  16. from ..utils.hook import HookHandler
  17. from ..utils.naming import AutoNaming
  18. logger = get_logger(__name__)
  19. def _expand_structure(prefix, obj):
  20. if isinstance(obj, (Tensor, Module)):
  21. return [(prefix, obj)]
  22. elif isinstance(obj, (list, tuple, dict)):
  23. ret = []
  24. if isinstance(obj, dict):
  25. targets = ((k, obj[k]) for k in sorted(obj))
  26. else:
  27. targets = ((str(k), v) for k, v in enumerate(obj))
  28. for k, o in targets:
  29. sub_ret = _expand_structure(k, o)
  30. if sub_ret and not isinstance(k, str):
  31. raise AssertionError(
  32. "keys for Tensor and Module must be str, error key: {}".format(k)
  33. )
  34. for kt, vt in sub_ret:
  35. ret.extend([(prefix + "." + kt, vt)])
  36. return ret
  37. else:
  38. return []
  39. def _access_structure(obj, key, callback=None):
  40. key_list = key.split(".")
  41. cur = obj
  42. parent = None
  43. for k in key_list:
  44. parent = cur
  45. if isinstance(cur, (list, tuple)):
  46. k = int(k)
  47. cur = cur[k]
  48. elif isinstance(cur, dict):
  49. cur = cur[k]
  50. else:
  51. cur = getattr(cur, k)
  52. return callback(parent, k, cur)
  53. def _is_parameter(obj):
  54. return isinstance(obj, Parameter)
  55. def _is_tensor(obj):
  56. return isinstance(obj, Tensor)
  57. def _is_buffer(obj):
  58. return isinstance(obj, Tensor) and not isinstance(obj, Parameter)
  59. def _is_module(obj):
  60. return isinstance(obj, Module)
  61. def _get_XNorm_typeclass():
  62. from .batchnorm import _BatchNorm
  63. from .normalization import GroupNorm, InstanceNorm, LayerNorm
  64. XNorm_types = (_BatchNorm, GroupNorm, LayerNorm, InstanceNorm)
  65. return XNorm_types
  66. class Module(metaclass=ABCMeta):
  67. r"""Base Module class.
  68. Args:
  69. name: module's name, can be initialized by the ``kwargs`` parameter
  70. of child class.
  71. """
  72. def __init__(self, name=None):
  73. self._modules = []
  74. if name is not None:
  75. assert (
  76. isinstance(name, str) and name.strip()
  77. ), "Module's name must be a non-empty string"
  78. self.name = name
  79. # runtime attributes
  80. self.training = True
  81. self.quantize_disabled = False
  82. # hooks
  83. self._forward_pre_hooks = OrderedDict()
  84. self._forward_hooks = OrderedDict()
  85. # used for profiler and automatic naming
  86. self._name = None
  87. self._short_name = None
  88. @abstractmethod
  89. def forward(self, inputs):
  90. pass
  91. def register_forward_pre_hook(self, hook: Callable) -> HookHandler:
  92. """Registers a hook to handle forward inputs. `hook` should be a function.
  93. Args:
  94. hook: a function that receive `module` and `inputs`, then return
  95. a modified `inputs` or `None`.
  96. Returns:
  97. a handler with :meth:`~.HookHandler.remove` interface to delete the hook.
  98. """
  99. return HookHandler(self._forward_pre_hooks, hook)
  100. def register_forward_hook(self, hook: Callable) -> HookHandler:
  101. """Registers a hook to handle forward results. `hook` should be a function that
  102. receive `module`, `inputs` and `outputs`, then return a modified `outputs` or `None`.
  103. This method return a handler with :meth:`~.HookHandler.remove` interface to delete the hook.
  104. """
  105. return HookHandler(self._forward_hooks, hook)
  106. def __call__(self, *inputs, **kwargs):
  107. AutoNaming.push_scope(self.name if self.name is not None else self._short_name)
  108. for hook in self._forward_pre_hooks.values():
  109. modified_inputs = hook(self, inputs)
  110. if modified_inputs is not None:
  111. if not isinstance(modified_inputs, tuple):
  112. modified_inputs = (modified_inputs,)
  113. inputs = modified_inputs
  114. outputs = self.forward(*inputs, **kwargs)
  115. for hook in self._forward_hooks.values():
  116. modified_outputs = hook(self, inputs, outputs)
  117. if modified_outputs is not None:
  118. outputs = modified_outputs
  119. AutoNaming.pop_scope()
  120. return outputs
  121. def _flatten(
  122. self,
  123. *,
  124. recursive: bool = True,
  125. with_key: bool = False,
  126. with_parent: bool = False,
  127. prefix: Optional[str] = None,
  128. predicate: Callable[[Any], bool] = lambda _: True,
  129. seen: Optional[Set[int]] = None
  130. ) -> Union[Iterable[Any], Iterable[Tuple[str, Any]]]:
  131. """Scans the module object and returns an iterable for the :class:`~.Tensor`
  132. and :class:`~.Module` attributes that agree with the ``predicate``. For multiple
  133. calls of this function with same arguments, the order of objects within the
  134. returned iterable is guaranteed to be identical, as long as all the involved
  135. module objects' ``__dict__`` does not change thoughout those calls.
  136. Args:
  137. recursive: whether to recursively scan all the submodules.
  138. with_key: whether to yield keys along with yielded objects.
  139. with_parent: whether to yield ``self`` along with yielded objects.
  140. prefix: prefix appended to the yielded keys.
  141. predicate: the predication function applied to scanned objects.
  142. seen: a dict that records whether a module has been traversed yet.
  143. """
  144. if seen is None:
  145. seen = set([id(self)])
  146. module_dict = vars(self)
  147. _prefix = "" if prefix is None else prefix + "."
  148. for key in sorted(module_dict):
  149. for expanded_key, leaf in _expand_structure(key, module_dict[key]):
  150. leaf_id = id(leaf)
  151. if leaf_id in seen:
  152. continue
  153. seen.add(leaf_id)
  154. if predicate(leaf):
  155. if with_key and with_parent:
  156. yield _prefix + expanded_key, leaf, self
  157. elif with_key:
  158. yield _prefix + expanded_key, leaf
  159. elif with_parent:
  160. yield leaf, self
  161. else:
  162. yield leaf
  163. if recursive and isinstance(leaf, Module):
  164. yield from leaf._flatten(
  165. recursive=recursive,
  166. with_key=with_key,
  167. with_parent=with_parent,
  168. prefix=_prefix + expanded_key if with_key else None,
  169. predicate=predicate,
  170. seen=seen,
  171. )
  172. def parameters(self, recursive: bool = True, **kwargs) -> Iterable[Parameter]:
  173. r"""Returns an iterable for the :class:`~.Parameter` of the module.
  174. Args:
  175. recursive: If ``True``, returns all :class:`~.Parameter` within this
  176. module, else only returns :class:`~.Parameter` that are direct attributes
  177. of this module.
  178. """
  179. if "requires_grad" in kwargs:
  180. del kwargs["requires_grad"]
  181. logger.warning(
  182. "Tensor currently has no requires_grad attribute "
  183. "so requires_grad argument is ignored here"
  184. )
  185. def predicate(obj) -> bool:
  186. return _is_parameter(obj)
  187. yield from self._flatten(
  188. with_key=False, predicate=predicate, recursive=recursive, **kwargs
  189. )
  190. def named_parameters(
  191. self, prefix: Optional[str] = None, recursive: bool = True, **kwargs
  192. ) -> Iterable[Tuple[str, Parameter]]:
  193. r"""Returns an iterable for key :class:`~.Parameter` pairs of the module, where
  194. ``key`` is the dotted path from this module to the :class:`~.Parameter`.
  195. Args:
  196. prefix: prefix prepended to the keys.
  197. recursive: if ``True``, returns all :class:`~.Parameter` within this
  198. module, else only returns :class:`~.Parameter` that are direct attributes
  199. of this module.
  200. """
  201. if "requires_grad" in kwargs:
  202. del kwargs["requires_grad"]
  203. logger.warning(
  204. "Tensor currently has no requires_grad attribute "
  205. "so requires_grad argument is ignored here"
  206. )
  207. def predicate(obj) -> bool:
  208. return _is_parameter(obj)
  209. yield from self._flatten(
  210. with_key=True,
  211. prefix=prefix,
  212. predicate=predicate,
  213. recursive=recursive,
  214. **kwargs,
  215. )
  216. def buffers(self, recursive: bool = True, **kwargs) -> Iterable[Tensor]:
  217. r"""Returns an iterable for the buffers of the module.
  218. Buffer is defined to be :class:`~.Tensor` excluding :class:`~.Parameter`.
  219. Args:
  220. recursive: if ``True``, returns all buffers within this
  221. module, else only returns buffers that are direct attributes
  222. """
  223. yield from self._flatten(
  224. with_key=False, predicate=_is_buffer, recursive=recursive, **kwargs
  225. )
  226. def named_buffers(
  227. self, prefix: Optional[str] = None, recursive: bool = True, **kwargs
  228. ) -> Iterable[Tuple[str, Tensor]]:
  229. r"""Returns an iterable for key buffer pairs of the module, where
  230. ``key`` is the dotted path from this module to the buffer.
  231. Buffer is defined to be :class:`~.Tensor` excluding :class:`~.Parameter`.
  232. Args:
  233. prefix: prefix prepended to the keys.
  234. recursive: if ``True``, returns all buffers within this
  235. module, else only returns buffers that are direct attributes
  236. of this module.
  237. prefix: Optional[str]:
  238. """
  239. yield from self._flatten(
  240. with_key=True,
  241. prefix=prefix,
  242. predicate=_is_buffer,
  243. recursive=recursive,
  244. **kwargs,
  245. )
  246. def tensors(self, recursive: bool = True, **kwargs) -> Iterable[Parameter]:
  247. r"""
  248. Returns an iterable for the :class:`~.Tensor` of the module.
  249. :param recursive: If ``True``, returns all :class:`~.Tensor` within this
  250. module, else only returns :class:`~.Tensor` that are direct attributes
  251. of this module.
  252. """
  253. yield from self._flatten(
  254. with_key=False, predicate=_is_tensor, recursive=recursive, **kwargs
  255. )
  256. def named_tensors(
  257. self, prefix: Optional[str] = None, recursive: bool = True, **kwargs
  258. ) -> Iterable[Tuple[str, Tensor]]:
  259. """
  260. Returns an iterable for key tensor pairs of the module, where
  261. ``key`` is the dotted path from this module to the tensor.
  262. :param prefix: prefix prepended to the keys.
  263. :param recursive: if ``True``, returns all tensors within this
  264. module, else only returns tensors that are direct attributes
  265. of this module.
  266. """
  267. yield from self._flatten(
  268. with_key=True,
  269. prefix=prefix,
  270. predicate=_is_tensor,
  271. recursive=recursive,
  272. **kwargs,
  273. )
  274. def children(self, **kwargs) -> "Iterable[Module]":
  275. r"""Returns an iterable for all the submodules that are direct attributes of this
  276. module.
  277. """
  278. yield from self._flatten(
  279. with_key=False, predicate=_is_module, recursive=False, **kwargs
  280. )
  281. def named_children(self, **kwargs) -> "Iterable[Tuple[str, Module]]":
  282. r"""Returns an iterable of key-submodule pairs for all the submodules that are
  283. direct attributes of this module, where 'key' is the attribute name of
  284. submodules.
  285. """
  286. yield from self._flatten(
  287. with_key=True, predicate=_is_module, recursive=False, **kwargs
  288. )
  289. def modules(self, **kwargs) -> "Iterable[Module]":
  290. r"""Returns an iterable for all the modules within this module, including itself."""
  291. if "with_parent" in kwargs and kwargs["with_parent"]:
  292. yield self, None
  293. else:
  294. yield self
  295. yield from self._flatten(with_key=False, predicate=_is_module, **kwargs)
  296. def named_modules(
  297. self, prefix: Optional[str] = None, **kwargs
  298. ) -> "Iterable[Tuple[str, Module]]":
  299. r"""Returns an iterable of key-module pairs for all the modules within this
  300. module, including itself, where 'key' is the dotted path from this module to the
  301. submodules.
  302. Args:
  303. prefix: prefix prepended to the path.
  304. """
  305. if "with_parent" in kwargs and kwargs["with_parent"]:
  306. yield ("" if prefix is None else prefix), self, None
  307. else:
  308. yield ("" if prefix is None else prefix), self
  309. yield from self._flatten(
  310. with_key=True, prefix=prefix, predicate=_is_module, **kwargs
  311. )
  312. def apply(self, fn: "Callable[[Module], Any]") -> None:
  313. r"""Applies function ``fn`` to all the modules within this module, including
  314. itself.
  315. Args:
  316. fn: the function to be applied on modules.
  317. """
  318. for it in self.modules():
  319. fn(it)
  320. @deprecated(version="1.0")
  321. def zero_grad(self) -> None:
  322. r"""Sets all parameters' grads to zero"""
  323. for param in self.parameters():
  324. if param.grad is not None:
  325. param.grad.reset_zero()
  326. def train(self, mode: bool = True, recursive: bool = True) -> None:
  327. r"""Sets training mode of all the modules within this module (including itself) to
  328. ``mode``. This effectively sets the ``training`` attributes of those modules
  329. to ``mode``, but only has effect on certain modules (e.g.
  330. :class:`~.BatchNorm2d`, :class:`~.Dropout`, :class:`~.Observer`)
  331. Args:
  332. mode: the training mode to be set on modules.
  333. recursive: whether to recursively call submodules' ``train()``.
  334. """
  335. if not recursive:
  336. self.training = mode
  337. return
  338. def fn(module: Module) -> None:
  339. module.train(mode, recursive=False)
  340. self.apply(fn)
  341. def eval(self) -> None:
  342. r"""Sets training mode of all the modules within this module (including itself) to
  343. ``False``. See :meth:`~.Module.train` for details.
  344. """
  345. self.train(False)
  346. def disable_quantize(self, value=True):
  347. r"""Sets ``module``'s ``quantize_disabled`` attribute and return ``module``.
  348. Could be used as a decorator.
  349. """
  350. def fn(module: Module) -> None:
  351. module.quantize_disabled = value
  352. self.apply(fn)
  353. @deprecated(version="1.0")
  354. def replace_param(
  355. self, params: dict, start_pos: int, seen: Optional[Set[int]] = None
  356. ):
  357. r"""Replaces module's parameters with ``params``, used by :class:`~.ParamPack` to
  358. speedup multimachine training.
  359. """
  360. offset = 0
  361. if seen is None:
  362. seen = set([id(self)])
  363. module_dict = vars(self)
  364. for key in sorted(module_dict):
  365. hash_id = id(module_dict[key])
  366. if hash_id in seen:
  367. continue
  368. seen.add(hash_id)
  369. if isinstance(module_dict[key], Parameter):
  370. if start_pos + offset in params:
  371. assert make_shape_tuple(module_dict[key].shape) == make_shape_tuple(
  372. params[start_pos + offset].shape
  373. )
  374. module_dict[key] = params[start_pos + offset]
  375. offset += 1
  376. if isinstance(module_dict[key], Module):
  377. offset += module_dict[key].replace_param(
  378. params, start_pos + offset, seen
  379. )
  380. return offset
  381. def state_dict(self, rst=None, prefix="", keep_var=False):
  382. _rst = self._state_dict(rst=rst, prefix=prefix, keep_var=keep_var)
  383. rst = OrderedDict()
  384. XNorm_typeclass = _get_XNorm_typeclass()
  385. for (module_type, k), v in _rst.items():
  386. # for performance reasons, parameters in XNorm (e.g., BatchNorm2d) are 4-dim tensors,
  387. # however they will be reshaped to 1-dim tensors before returned by `statr_dict()`
  388. if issubclass(module_type, XNorm_typeclass):
  389. v = v.reshape(-1)
  390. rst[k] = v
  391. return rst
  392. def _state_dict(self, rst=None, prefix="", keep_var=False):
  393. r"""Returns a dictionary containing whole states of the module."""
  394. def is_state(obj):
  395. return _is_parameter(obj) or _is_buffer(obj)
  396. module_type = self.__class__
  397. if rst is None:
  398. rst = OrderedDict()
  399. for k, v in self._flatten(recursive=False, with_key=True, predicate=is_state):
  400. assert prefix + k not in rst, "duplicated state: {}".format(k)
  401. if keep_var:
  402. rst[(module_type, prefix + k)] = v
  403. else:
  404. rst[(module_type, prefix + k)] = v.numpy()
  405. for k, submodule in self._flatten(
  406. recursive=False,
  407. with_key=True,
  408. predicate=lambda obj: isinstance(obj, Module),
  409. ):
  410. submodule.state_dict(rst, prefix + k + ".", keep_var)
  411. return rst
  412. def load_state_dict(
  413. self,
  414. state_dict: Union[dict, Callable[[str, Tensor], Optional[np.ndarray]]],
  415. strict=True,
  416. ):
  417. r"""Loads a given dictionary created by :func:`state_dict` into this module.
  418. If ``strict`` is ``True``, the keys of :func:`state_dict` must exactly match the keys
  419. returned by :func:`state_dict`.
  420. Users can also pass a closure: ``Function[key: str, var: Tensor] -> Optional[np.ndarray]``
  421. as a `state_dict`, in order to handle complex situations. For example, load everything
  422. except for the final linear classifier:
  423. .. code-block::
  424. state_dict = {...} # Dict[str, np.ndarray]
  425. model.load_state_dict({
  426. k: None if k.startswith('fc') else v
  427. for k, v in state_dict.items()
  428. }, strict=False)
  429. Here returning ``None`` means skipping parameter ``k``.
  430. To prevent shape mismatch (e.g. load PyTorch weights), we can reshape before loading:
  431. .. code-block::
  432. state_dict = {...}
  433. def reshape_accordingly(k, v):
  434. return state_dict[k].reshape(v.shape)
  435. model.load_state_dict(reshape_accordingly)
  436. We can also perform inplace re-initialization or pruning:
  437. .. code-block::
  438. def reinit_and_pruning(k, v):
  439. if 'bias' in k:
  440. M.init.zero_(v)
  441. if 'conv' in k:
  442. """
  443. unused = []
  444. if isinstance(state_dict, dict):
  445. unused = state_dict.keys()
  446. def closure(k, _): # var unused
  447. return state_dict[k] if k in state_dict else None
  448. elif callable(state_dict):
  449. closure = state_dict
  450. else:
  451. raise ValueError(
  452. "`state_dict` must load a dict or callable, got {}".format(
  453. type(state_dict)
  454. )
  455. )
  456. loaded, skipped = self._load_state_dict_with_closure(closure)
  457. unused = set(unused) - loaded
  458. if len(unused) != 0:
  459. if strict:
  460. raise KeyError(
  461. "Unused params violate `strict=True`, unused={}".format(unused)
  462. )
  463. else:
  464. logger.warning(
  465. "Unused params in `strict=False` mode, unused={}".format(unused)
  466. )
  467. if len(skipped) != 0:
  468. if strict:
  469. raise KeyError(
  470. "Missing params violate `strict=True`, missing={}".format(skipped)
  471. )
  472. else:
  473. logger.warning(
  474. "Missing params in `strict=False` mode, missing={}".format(skipped)
  475. )
  476. def _load_state_dict_with_closure(self, closure):
  477. r"""Advance state_dict load through callable ``closure`` whose signature is
  478. ``closure(key: str, var: Tensor) -> Union[np.ndarry, None]``
  479. """
  480. XNorm_typeclass = _get_XNorm_typeclass()
  481. assert callable(closure), "closure must be a function"
  482. loaded = []
  483. skipped = []
  484. local_state_dict = self._state_dict(keep_var=True)
  485. for (module_type, k), var in local_state_dict.items():
  486. to_be_load = closure(k, var)
  487. if to_be_load is None:
  488. skipped.append(k)
  489. continue
  490. assert isinstance(
  491. to_be_load, np.ndarray
  492. ), "closure should return a `np.ndarray`, now `{}` get {}".format(
  493. k, to_be_load
  494. )
  495. var_shape = make_shape_tuple(var.shape)
  496. to_be_load_shape = make_shape_tuple(to_be_load.shape)
  497. if var_shape != to_be_load_shape:
  498. # weight and bias in BatchNorm1d, BatchNorm2d and SyncBatchNorm are 1-dim tensors in v1.0, and
  499. # since v1.1 they are 4-dim tensors. The following special rule for these modules preserves the
  500. # backward compatibility.
  501. if issubclass(module_type, XNorm_typeclass):
  502. if np.prod(var_shape) == np.prod(to_be_load_shape):
  503. to_be_load = to_be_load.reshape(var_shape)
  504. else:
  505. raise ValueError(
  506. "param `{}` size mismatch, should be {}, get {}".format(
  507. k, np.prod(var_shape), np.prod(to_be_load_shape)
  508. )
  509. )
  510. else:
  511. raise ValueError(
  512. "param `{}` shape mismatch, should be {}, get {}".format(
  513. k, var_shape, to_be_load_shape
  514. )
  515. )
  516. var._reset(
  517. type(var)(
  518. to_be_load, dtype=to_be_load.dtype, device=var.device, no_cache=True
  519. )
  520. )
  521. loaded.append(k)
  522. return set(loaded), set(skipped)
  523. def __setattr__(self, name: str, value):
  524. is_module_like = _is_module(value) or isinstance(value, (list, tuple, dict))
  525. if name != "_modules":
  526. modules = self.__dict__.get("_modules")
  527. if modules is None and is_module_like:
  528. raise AttributeError(
  529. "cannot assign module before Module.__init__() call"
  530. )
  531. if is_module_like:
  532. if name not in modules:
  533. modules.append(name)
  534. else:
  535. if modules is not None and name in modules:
  536. modules.remove(name)
  537. def append_name(prefix, name):
  538. if prefix is None or prefix == "":
  539. return name
  540. return prefix + "." + name
  541. def set_name(parent, prefix, name, obj):
  542. if isinstance(obj, Tensor):
  543. assert obj.name is not None
  544. if obj.name != "":
  545. name = obj.name
  546. full_name = append_name(prefix, name)
  547. if obj._short_name and obj._short_name != name:
  548. logger.warning(
  549. "try setting the submodule `{}` to `{}`'s new attribute `{}`, its name `{}` will remain unchanged".format(
  550. obj._short_name, type(parent), name, obj._short_name
  551. )
  552. )
  553. return
  554. if isinstance(obj, Tensor):
  555. obj._prefix = prefix
  556. obj._name = full_name
  557. obj._short_name = name
  558. obj._set_name(obj._name)
  559. return obj._name
  560. elif isinstance(obj, Module):
  561. obj._name = full_name
  562. obj._short_name = name
  563. for k, v in obj._flatten(recursive=False, with_key=True):
  564. set_name(obj, full_name, k, v)
  565. return obj._name
  566. else:
  567. assert False
  568. for k, v in _expand_structure(name, value):
  569. prefix = self._name if self._name else self.name
  570. set_name(self, prefix, k, v)
  571. super().__setattr__(name, value)
  572. def __delattr__(self, name: str):
  573. if name in self.__dict__ and _is_module(self.__dict__[name]):
  574. modules = self.__dict__.get("_modules")
  575. if name in modules:
  576. modules.remove(name)
  577. super().__delattr__(name)
  578. def _module_info_string(self) -> str:
  579. r"""Set the extra representation of the module."""
  580. return ""
  581. def __repr__(self):
  582. def add_indent(repr_str, num_spaces):
  583. s = repr_str.split("\n")
  584. # don't do anything for single-line stuff
  585. if len(s) == 1:
  586. return repr_str
  587. first = s.pop(0)
  588. s = [(num_spaces * " ") + line for line in s]
  589. s = "\n".join(s)
  590. s = first + "\n" + s
  591. return s
  592. extra_lines = []
  593. extra_repr = self._module_info_string()
  594. if extra_repr:
  595. extra_lines = extra_repr.split("\n")
  596. child_lines = []
  597. for name in self._modules:
  598. if _is_module(self.__dict__[name]):
  599. child_lines.append(
  600. "(" + name + "): " + add_indent(repr(self.__dict__[name]), 2)
  601. )
  602. else:
  603. for k, v in _expand_structure(name, self.__dict__[name]):
  604. if _is_module(v):
  605. child_lines.append("(" + k + "): " + add_indent(repr(v), 2))
  606. lines = extra_lines + child_lines
  607. main_str = self.__class__.__name__ + "("
  608. if lines:
  609. # simple one-liner info, which most builtin Modules will use
  610. if len(extra_lines) == 1 and not child_lines:
  611. main_str += extra_lines[0]
  612. else:
  613. main_str += "\n " + "\n ".join(lines) + "\n"
  614. main_str += ")"
  615. return main_str