fp8.py 5.4 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152
  1. from typing import Any, Dict, List, Optional
  2. import torch
  3. from torch.nn import Module
  4. from torch.nn.parameter import Parameter
  5. from aphrodite.modeling.layers.linear import (LinearMethodBase,
  6. set_weight_attrs)
  7. from aphrodite.quantization.base_config import QuantizationConfig
  8. class FP8Config(QuantizationConfig):
  9. """Config class for FP8."""
  10. @classmethod
  11. def get_name(cls) -> str:
  12. return "fp8"
  13. @classmethod
  14. def get_supported_act_dtypes(cls) -> List[torch.dtype]:
  15. return [torch.bfloat16, torch.half]
  16. @classmethod
  17. def get_min_capability(cls) -> int:
  18. return 89
  19. @classmethod
  20. def get_config_filenames(cls) -> List[str]:
  21. return []
  22. @classmethod
  23. def from_config(cls, config: Dict[str, Any]) -> "FP8Config":
  24. return cls()
  25. def get_linear_method(self) -> "Fp8LinearMethod":
  26. return Fp8LinearMethod(self)
  27. def get_scaled_act_names(self) -> List[str]:
  28. return []
  29. def merge_weight(self) -> bool:
  30. return True
  31. def rope_style(self) -> Optional[bool]:
  32. return None
  33. def quant_vocab(self) -> List[bool]:
  34. return [False, False]
  35. def support_fused_moe(self) -> bool:
  36. return True
  37. class Fp8LinearMethod(LinearMethodBase):
  38. """Linear method for FP8.
  39. We now support common FP16/BF16 model checkpoints ONLY. The weight
  40. scaling factor will be initialized after the model weights are loaded.
  41. Limitations:
  42. 1. Only support per-tensor quantization due to torch._scaled_mm support.
  43. 2. Only support float8_e4m3fn data type due to the limitation of
  44. torch._scaled_mm (https://github.com/pytorch/pytorch/blob/2e48b39603411a41c5025efbe52f89560b827825/aten/src/ATen/native/cuda/Blas.cpp#L854-L856)
  45. Args:
  46. quant_config: The quantization config.
  47. """
  48. def __init__(self, quant_config: FP8Config):
  49. self.quant_config = quant_config
  50. def create_weights(
  51. self,
  52. layer: torch.nn.Module,
  53. input_size_per_partition: int,
  54. output_partition_sizes: List[int],
  55. input_size: int,
  56. output_size: int,
  57. params_dtype: torch.dtype,
  58. **extra_weight_attrs,
  59. ):
  60. output_size_per_partition = sum(output_partition_sizes)
  61. weight = Parameter(torch.empty(output_size_per_partition,
  62. input_size_per_partition,
  63. dtype=params_dtype),
  64. requires_grad=False)
  65. layer.register_parameter("weight", weight)
  66. set_weight_attrs(weight, {"input_dim": 1, "output_dim": 0})
  67. set_weight_attrs(weight, extra_weight_attrs)
  68. w_scale = Parameter(
  69. torch.empty(1, dtype=torch.float32),
  70. requires_grad=False,
  71. )
  72. layer.register_parameter("weight_scaling_factor", w_scale)
  73. def process_weights_after_loading(self, layer: Module) -> None:
  74. # Although the linear_method is propagated to all layers,
  75. # only linear layers invoke "create_weights". So we check
  76. # whether "weight_scaling_facor" is registered to determine
  77. # whether the layer is a linear layer that requires quantization.
  78. if not hasattr(layer, "weight_scaling_factor"):
  79. return
  80. qweight, weight_scale = per_tensor_quantize(layer.weight)
  81. # torch._scaled_mm requires column-major in the second
  82. # input (weight), so we transpose the quantized weight.
  83. layer.weight = Parameter(qweight.t(), requires_grad=False)
  84. layer.weight_scaling_factor.data.copy_(weight_scale)
  85. def apply_weights(self,
  86. layer: torch.nn.Module,
  87. x: torch.Tensor,
  88. bias: Optional[torch.Tensor] = None) -> torch.Tensor:
  89. qinput, x_scale = per_tensor_quantize(x)
  90. output, _ = torch._scaled_mm(
  91. qinput,
  92. layer.weight,
  93. out_dtype=x.dtype,
  94. scale_a=x_scale,
  95. scale_b=layer.weight_scaling_factor,
  96. bias=bias,
  97. )
  98. return output
  99. def apply_moe_weights(self, w1: Dict[str,
  100. torch.Tensor], w2: Dict[str,
  101. torch.Tensor],
  102. x: torch.Tensor, gating_output: torch.Tensor,
  103. topk: int, renormalize: bool) -> torch.Tensor:
  104. raise NotImplementedError
  105. def per_tensor_quantize(tensor: torch.Tensor) -> tuple[torch.Tensor, float]:
  106. """Quantize a tensor using per-tensor static scaling factor.
  107. Args:
  108. tensor: The input tensor.
  109. """
  110. finfo = torch.finfo(torch.float8_e4m3fn)
  111. # Calculate the scale as dtype max divided by absmax.
  112. # Since .abs() creates a new tensor, we use aminmax to get
  113. # the min and max first and then calculate the absmax.
  114. min_val, max_val = tensor.aminmax()
  115. amax = min_val.abs().max(max_val.abs())
  116. scale = finfo.max / amax.clamp(min=1e-12)
  117. # scale and clamp the tensor to bring it to
  118. # the representative range of float8 data type
  119. # (as default cast is unsaturated)
  120. qweight = (tensor * scale).clamp(min=finfo.min, max=finfo.max)
  121. # Return both float8 data and the inverse scale (as float),
  122. # as both required as inputs to torch._scaled_mm
  123. qweight = qweight.to(torch.float8_e4m3fn)
  124. scale = scale.float().reciprocal()
  125. return qweight, scale