lazy.py 11 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261
  1. import itertools
  2. from typing_extensions import Protocol
  3. import warnings
  4. import torch
  5. from ..parameter import is_lazy
  6. class _LazyProtocol(Protocol):
  7. """This is to avoid errors with mypy checks for
  8. The attributes in a mixin:
  9. https://mypy.readthedocs.io/en/latest/more_types.html#mixin-classes
  10. """
  11. def _register_load_state_dict_pre_hook(self, hook):
  12. ...
  13. def register_forward_pre_hook(self, hook):
  14. ...
  15. def _lazy_load_hook(
  16. self, state_dict, prefix, local_metadata, strict,
  17. missing_keys, unexpected_keys, error_msgs):
  18. ...
  19. def _get_name(self):
  20. ...
  21. def _infer_parameters(self, module, input):
  22. ...
  23. @property
  24. def _parameters(self):
  25. ...
  26. @property
  27. def _buffers(self):
  28. ...
  29. @property
  30. def _non_persistent_buffers_set(self):
  31. ...
  32. @property
  33. def _load_hook(self):
  34. ...
  35. @property
  36. def _initialize_hook(self):
  37. ...
  38. class LazyModuleMixin:
  39. r"""A mixin for modules that lazily initialize parameters, also known as "lazy modules."
  40. .. warning:
  41. Lazy modules are an experimental new feature under active development,
  42. and their API is likely to change.
  43. Modules that lazily initialize parameters, or "lazy modules",
  44. derive the shapes of their parameters from the first input(s)
  45. to their forward method. Until that first forward they contain
  46. :class:`torch.nn.UninitializedParameter` s that should not be accessed
  47. or used, and afterward they contain regular :class:`torch.nn.Parameter` s.
  48. Lazy modules are convenient since they don't require computing some
  49. module arguments, like the :attr:`in_features` argument of a
  50. typical :class:`torch.nn.Linear`.
  51. After construction, networks with lazy modules should first
  52. be converted to the desired dtype and placed on the expected device.
  53. This is because lazy modules only perform shape inference so the usual dtype
  54. and device placement behavior applies.
  55. The lazy modules should then perform "dry runs" to initialize all the components in the module.
  56. These "dry runs" send inputs of the correct size, dtype, and device through
  57. the network and to each one of its lazy modules. After this the network can be used as usual.
  58. >>> class LazyMLP(torch.nn.Module):
  59. ... def __init__(self):
  60. ... super().__init__()
  61. ... self.fc1 = torch.nn.LazyLinear(10)
  62. ... self.relu1 = torch.nn.ReLU()
  63. ... self.fc2 = torch.nn.LazyLinear(1)
  64. ... self.relu2 = torch.nn.ReLU()
  65. ...
  66. ... def forward(self, input):
  67. ... x = self.relu1(self.fc1(input))
  68. ... y = self.relu2(self.fc2(x))
  69. ... return y
  70. >>> # constructs a network with lazy modules
  71. >>> lazy_mlp = LazyMLP()
  72. >>> # transforms the network's device and dtype
  73. >>> # NOTE: these transforms can and should be applied after construction and before any 'dry runs'
  74. >>> lazy_mlp = lazy_mlp.cuda().double()
  75. >>> lazy_mlp
  76. LazyMLP( (fc1): LazyLinear(in_features=0, out_features=10, bias=True)
  77. (relu1): ReLU()
  78. (fc2): LazyLinear(in_features=0, out_features=1, bias=True)
  79. (relu2): ReLU()
  80. )
  81. >>> # performs a dry run to initialize the network's lazy modules
  82. >>> lazy_mlp(torch.ones(10,10).cuda())
  83. >>> # after initialization, LazyLinear modules become regular Linear modules
  84. >>> lazy_mlp
  85. LazyMLP(
  86. (fc1): Linear(in_features=10, out_features=10, bias=True)
  87. (relu1): ReLU()
  88. (fc2): Linear(in_features=10, out_features=1, bias=True)
  89. (relu2): ReLU()
  90. )
  91. >>> # attaches an optimizer, since parameters can now be used as usual
  92. >>> optim = torch.optim.SGD(mlp.parameters(), lr=0.01)
  93. A final caveat when using lazy modules is that the order of initialization of a network's
  94. parameters may change, since the lazy modules are always initialized after other modules.
  95. For example, if the LazyMLP class defined above had a :class:`torch.nn.LazyLinear` module
  96. first and then a regular :class:`torch.nn.Linear` second, the second module would be
  97. initialized on construction and the first module would be initialized during the first dry run.
  98. This can cause the parameters of a network using lazy modules to be initialized differently
  99. than the parameters of a network without lazy modules as the order of parameter initializations,
  100. which often depends on a stateful random number generator, is different.
  101. Check :doc:`/notes/randomness` for more details.
  102. Lazy modules can be serialized with a state dict like other modules. For example:
  103. >>> lazy_mlp = LazyMLP()
  104. >>> # The state dict shows the uninitialized parameters
  105. >>> lazy_mlp.state_dict()
  106. OrderedDict([('fc1.weight', Uninitialized parameter),
  107. ('fc1.bias',
  108. tensor([-1.8832e+25, 4.5636e-41, -1.8832e+25, 4.5636e-41, -6.1598e-30,
  109. 4.5637e-41, -1.8788e+22, 4.5636e-41, -2.0042e-31, 4.5637e-41])),
  110. ('fc2.weight', Uninitialized parameter),
  111. ('fc2.bias', tensor([0.0019]))])
  112. Lazy modules can load regular :class:`torch.nn.Parameter` s (i.e. you can serialize/deserialize
  113. initialized LazyModules and they will remain initialized)
  114. >>> full_mlp = LazyMLP()
  115. >>> # Dry run to initialize another module
  116. >>> full_mlp.forward(torch.ones(10, 1))
  117. >>> # Load an initialized state into a lazy module
  118. >>> lazy_mlp.load_state_dict(full_mlp.state_dict())
  119. >>> # The state dict now holds valid values
  120. >>> lazy_mlp.state_dict()
  121. OrderedDict([('fc1.weight',
  122. tensor([[-0.3837],
  123. [ 0.0907],
  124. [ 0.6708],
  125. [-0.5223],
  126. [-0.9028],
  127. [ 0.2851],
  128. [-0.4537],
  129. [ 0.6813],
  130. [ 0.5766],
  131. [-0.8678]])),
  132. ('fc1.bias',
  133. tensor([-1.8832e+25, 4.5636e-41, -1.8832e+25, 4.5636e-41, -6.1598e-30,
  134. 4.5637e-41, -1.8788e+22, 4.5636e-41, -2.0042e-31, 4.5637e-41])),
  135. ('fc2.weight',
  136. tensor([[ 0.1320, 0.2938, 0.0679, 0.2793, 0.1088, -0.1795, -0.2301, 0.2807,
  137. 0.2479, 0.1091]])),
  138. ('fc2.bias', tensor([0.0019]))])
  139. Note, however, that the loaded parameters will not be replaced when doing a "dry run" if they are initialized
  140. when the state is loaded. This prevents using initialized modules in different contexts.
  141. """
  142. # modules inheriting from this will change their __class__ to the specified
  143. # one after they are fully initialized
  144. cls_to_become = None
  145. def __init__(self: _LazyProtocol, *args, **kwargs):
  146. # Mypy doesnt like this super call in a mixin
  147. super().__init__(*args, **kwargs) # type: ignore[misc]
  148. self._load_hook = self._register_load_state_dict_pre_hook(self._lazy_load_hook)
  149. self._initialize_hook = self.register_forward_pre_hook(self._infer_parameters)
  150. warnings.warn('Lazy modules are a new feature under heavy development '
  151. 'so changes to the API or functionality can happen at any moment.')
  152. def _save_to_state_dict(self: _LazyProtocol, destination, prefix, keep_vars):
  153. # This should be ideally implemented as a hook,
  154. # but we should override `detach` in the UninitializedParameter to return itself
  155. # which is not clean
  156. for name, param in self._parameters.items():
  157. if param is not None:
  158. if not (is_lazy(param) or keep_vars):
  159. param = param.detach()
  160. destination[prefix + name] = param
  161. for name, buf in self._buffers.items():
  162. if buf is not None and name not in self._non_persistent_buffers_set:
  163. if not (is_lazy(buf) or keep_vars):
  164. buf = buf.detach()
  165. destination[prefix + name] = buf
  166. def _lazy_load_hook(
  167. self: _LazyProtocol, state_dict, prefix, local_metadata, strict,
  168. missing_keys, unexpected_keys, error_msgs):
  169. """load_state_dict pre-hook function for lazy buffers and parameters.
  170. The purpose of this hook is to adjust the current state and/or
  171. ``state_dict`` being loaded so that a module instance serialized in
  172. both un/initialized state can be deserialized onto both un/initialized
  173. module instance.
  174. See comment in ``torch.nn.Module._register_load_state_dict_pre_hook``
  175. for the details of the hook specification.
  176. """
  177. for name, param in itertools.chain(self._parameters.items(), self._buffers.items()):
  178. key = prefix + name
  179. if key in state_dict and param is not None:
  180. input_param = state_dict[key]
  181. if is_lazy(param):
  182. # The current parameter is not initialized but the one being loaded one is
  183. # create a new parameter based on the uninitialized one
  184. if not is_lazy(input_param):
  185. with torch.no_grad():
  186. param.materialize(input_param.shape)
  187. def initialize_parameters(self: _LazyProtocol, *args, **kwargs):
  188. r"""Initialize parameters according to the input batch properties.
  189. This adds an interface to isolate parameter initialization from the
  190. forward pass when doing parameter shape inference.
  191. """
  192. raise NotImplementedError('initialize_parameters is not implemented for {}'.format(self.__class__.__name__))
  193. def has_uninitialized_params(self: _LazyProtocol):
  194. r"""Check if a module has parameters that are not initialized
  195. """
  196. # This is to avoid the JIT to track this parameter and force
  197. # custom modules __setstate__ to add it
  198. params = self._parameters.values()
  199. buffers = self._buffers.values()
  200. for param in itertools.chain(params, buffers):
  201. if is_lazy(param):
  202. return True
  203. return False
  204. def _infer_parameters(self: _LazyProtocol, module, input):
  205. r"""Infers the size and initializes the parameters according to the
  206. provided input batch.
  207. Given a module that contains parameters that were declared inferrable
  208. using :class:`torch.nn.parameter.ParameterMode.Infer`, runs a forward pass
  209. in the complete module using the provided input to initialize all the parameters
  210. as needed.
  211. The module is set into evaluation mode before running the forward pass in order
  212. to avoid saving statistics or calculating gradients
  213. """
  214. module.initialize_parameters(*input)
  215. if module.has_uninitialized_params():
  216. raise RuntimeError('module {} has not been fully initialized'.format(self._get_name()))
  217. module._initialize_hook.remove()
  218. module._load_hook.remove()
  219. delattr(module, '_initialize_hook')
  220. delattr(module, '_load_hook')
  221. if module.cls_to_become is not None:
  222. module.__class__ = module.cls_to_become
  223. def _replicate_for_data_parallel(self: _LazyProtocol):
  224. raise RuntimeError('Modules with uninitialized parameters can\'t be used with `DataParallel`. '
  225. 'Run a dummy forward pass to correctly initialize the modules')