eigenvalue.py 5.7 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154
  1. '''Copyright The Microsoft DeepSpeed Team'''
  2. import torch
  3. from deepspeed.utils import log_dist
  4. import numpy as np
  5. import logging
  6. class Eigenvalue(object):
  7. def __init__(self,
  8. verbose=False,
  9. max_iter=100,
  10. tol=1e-2,
  11. stability=0,
  12. gas_boundary_resolution=1,
  13. layer_name='',
  14. layer_num=0):
  15. super().__init__()
  16. self.verbose = verbose
  17. self.max_iter = max_iter
  18. self.tol = tol
  19. self.stability = stability
  20. self.gas_boundary_resolution = gas_boundary_resolution
  21. self.layer_name = layer_name
  22. self.layer_num = layer_num
  23. assert len(self.layer_name) > 0 and layer_num > 0
  24. log_dist(
  25. f'enabled eigenvalue with verbose={verbose}, max_iter={max_iter}, tol={tol}, stability={stability}, gas_boundary_resolution={gas_boundary_resolution}, layer_name={layer_name}, layer_num={layer_num}',
  26. ranks=[0])
  27. # Replace all nan/pos-inf/neg-inf to zero
  28. # TODO: Pytorch new version may add this function, replace this one by then.
  29. def nan_to_num(self, x):
  30. device = x.device
  31. x = x.cpu().numpy()
  32. x = np.nan_to_num(x=x, copy=False, nan=0.0, posinf=0.0, neginf=0.0)
  33. return torch.from_numpy(x).to(device)
  34. def normalize(self, v):
  35. norm_squared = self.inner_product(v, v)
  36. norm = norm_squared**0.5 + self.stability
  37. normalized_vectors = [vector / norm for vector in v]
  38. normalized_vectors = [self.nan_to_num(vector) for vector in normalized_vectors]
  39. return normalized_vectors
  40. def inner_product(self, xs, ys):
  41. return sum([torch.sum(x * y) for (x, y) in zip(xs, ys)])
  42. def get_layers(self, module):
  43. scope_names = self.layer_name.split('.')
  44. assert len(scope_names) > 0
  45. m = module
  46. for name in scope_names:
  47. assert hasattr(m, name), "layer_name configuration is invalid."
  48. m = getattr(m, name)
  49. return m
  50. def compute_eigenvalue(self, module, device=None, scale=1.0):
  51. block_eigenvalue = []
  52. param_keys = []
  53. layers = self.get_layers(module)
  54. for block in range(self.layer_num):
  55. model_block = layers[block]
  56. # We found this randn() has obvious accuracy impact in some cases, save/recover random state here.
  57. rng_state = torch.random.get_rng_state()
  58. if device is None:
  59. v = [
  60. torch.randn(p.size()) for p in model_block.parameters()
  61. if p.grad is not None and p.grad.grad_fn is not None
  62. ]
  63. else:
  64. v = [
  65. torch.randn(p.size(),
  66. device=device) for p in model_block.parameters()
  67. if p.grad is not None and p.grad.grad_fn is not None
  68. ]
  69. torch.random.set_rng_state(rng_state)
  70. grads = [
  71. param.grad for param in model_block.parameters()
  72. if param.grad is not None and param.grad.grad_fn is not None
  73. ]
  74. params = [
  75. param for param in model_block.parameters()
  76. if param.grad is not None and param.grad.grad_fn is not None
  77. ]
  78. layer_keys = [id(p) for p in model_block.parameters()]
  79. param_keys.append(layer_keys)
  80. v = self.normalize(v)
  81. # Disable eigenvalue if the model doesn't support second order gradients computation,
  82. # e.g. when enabling DS transformer kernel.
  83. if len(grads) == 0 or len(params) == 0:
  84. log_dist(f'The model does NOT support eigenvalue computation.',
  85. ranks=[0],
  86. level=logging.WARNING)
  87. return []
  88. i = 0
  89. eigenvalue_current, eigenvalue_previous = 1., 0.
  90. while (i < self.max_iter) and abs(eigenvalue_current) > 0 and (abs(
  91. (eigenvalue_current - eigenvalue_previous) /
  92. eigenvalue_current) >= self.tol): # test convergence criteria
  93. eigenvalue_previous = eigenvalue_current
  94. Hv = torch.autograd.grad(grads,
  95. params,
  96. grad_outputs=v,
  97. only_inputs=True,
  98. retain_graph=True)
  99. #Hv = [hv.float() for hv in Hv]
  100. Hv = [self.nan_to_num(hv).float() for hv in Hv]
  101. eigenvalue_current = self.inner_product(Hv, v).item()
  102. v = self.normalize(Hv)
  103. v = [x / scale for x in v]
  104. i += 1
  105. eigenvalue_current *= scale
  106. block_eigenvalue.append(eigenvalue_current)
  107. if self.verbose:
  108. log_dist(
  109. f'block: {block}, power iteration: {i}, eigenvalue: {eigenvalue_current}',
  110. ranks=[0])
  111. block_eigenvalue = self.post_process(block_eigenvalue)
  112. if self.verbose:
  113. log_dist(f'post processed block_eigenvalue: {block_eigenvalue}', ranks=[0])
  114. # {param_id: (eigenvalue, layer_id)}
  115. ev_dict = {}
  116. for i, (layer_keys, value) in enumerate(zip(param_keys, block_eigenvalue)):
  117. ev_dict.update(dict.fromkeys(layer_keys, (value, i)))
  118. return ev_dict
  119. # 1. Map all eigenvalues to [0, 1.0].
  120. # 2. Some layers can't generate valid eigenvalues on fp16 precision, use 1.0 instead.
  121. def post_process(self, value_list):
  122. max_value = abs(max(value_list, key=abs))
  123. return [abs(v) / max_value if v != 0.0 else 1.0 for v in value_list]