exl2.py 5.2 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159
  1. from contextlib import suppress
  2. from typing import Any, Dict, List, Optional
  3. import torch
  4. from aphrodite.modeling.layers.linear import LinearBase, LinearMethodBase
  5. from aphrodite.modeling.utils import set_weight_attrs
  6. from aphrodite.quantization.base_config import QuantizationConfig
  7. HAS_QUANTS = False
  8. with suppress(ImportError):
  9. from aphrodite._quant_C import quant_ops as ops
  10. HAS_QUANTS = True
  11. def make_group_map(q_groups, num_qrows):
  12. gr = q_groups.tolist()
  13. group_map = []
  14. num_groups = len(gr) // 2
  15. for i in range(num_groups):
  16. bits = gr[i * 2]
  17. if i < num_groups - 1:
  18. qrows = gr[i * 2 + 3] - gr[i * 2 + 1]
  19. else:
  20. qrows = num_qrows - gr[i * 2 + 1]
  21. rows = qrows * 32 // bits
  22. for j in range(rows):
  23. group_map += [i]
  24. group_map += [rows - j]
  25. return torch.tensor(group_map, dtype=torch.short, device=q_groups.device)
  26. class Exl2Config(QuantizationConfig):
  27. """Config class for Exl2."""
  28. def __repr__(self) -> str:
  29. return "Exl2Config()"
  30. @classmethod
  31. def get_name(cls) -> str:
  32. return "exl2"
  33. @classmethod
  34. def get_supported_act_dtypes(cls) -> List[torch.dtype]:
  35. return [torch.half]
  36. @classmethod
  37. # Need to figure it out
  38. def get_min_capability(cls) -> int:
  39. return 60
  40. @classmethod
  41. def get_config_filenames(cls) -> List[str]:
  42. return []
  43. @classmethod
  44. def from_config(cls, config: Dict[str, Any]) -> "Exl2Config":
  45. return cls()
  46. def get_quant_method(
  47. self, layer: torch.nn.Module) -> Optional["Exl2LinearMethod"]:
  48. if isinstance(layer, LinearBase):
  49. return Exl2LinearMethod(self)
  50. return None
  51. def get_scaled_act_names(self) -> List[str]:
  52. return []
  53. def merge_weight(self) -> bool:
  54. return False
  55. def quant_vocab(self) -> List[bool]:
  56. return [False, True]
  57. def support_fused_moe(self) -> bool:
  58. return False
  59. def rope_style(self) -> Optional[bool]:
  60. return None
  61. class Exl2LinearMethod(LinearMethodBase):
  62. """Linear method for Exl2.
  63. Args:
  64. quant_config: The Exl2 quantization config.
  65. """
  66. def __init__(self, quant_config: Exl2Config):
  67. if not HAS_QUANTS:
  68. raise ImportError("Could not find the quantization kernels.")
  69. self.quant_config = quant_config
  70. def create_weights(self, layer: torch.nn.Module,
  71. input_size_per_partition: int,
  72. output_partition_sizes: List[int], input_size: int,
  73. output_size: int, params_dtype: torch.dtype,
  74. **extra_weight_attr):
  75. # The shape of weight is unknown until load state dict
  76. # q_groups, q_invperm, q_scale, q_scale_max, q_weight, q_groups
  77. layer.exllama_state = 0
  78. qweight = torch.nn.parameter.UninitializedParameter(
  79. requires_grad=False)
  80. set_weight_attrs(qweight, {"output_dim": 1, "ignore_warning": True})
  81. layer.register_parameter("q_weight", qweight)
  82. qscale = torch.nn.parameter.UninitializedParameter(requires_grad=False)
  83. set_weight_attrs(
  84. qscale, {
  85. "output_dim": 1,
  86. "packed_dim": 1,
  87. "pack_factor": 8,
  88. "ignore_warning": True
  89. })
  90. layer.register_parameter("q_scale", qscale)
  91. for name in ["q_groups", "q_invperm", "q_scale_max"]:
  92. fake_weight = torch.nn.parameter.UninitializedParameter(
  93. requires_grad=False)
  94. set_weight_attrs(fake_weight, {"ignore_warning": True})
  95. layer.register_parameter(name, fake_weight)
  96. def apply(self,
  97. layer: torch.nn.Module,
  98. x: torch.Tensor,
  99. bias: Optional[torch.Tensor] = None) -> torch.Tensor:
  100. out_shape = x.shape[:-1] + (layer.q_weight.shape[-1], )
  101. reshaped_x = x.reshape(-1, x.shape[-1])
  102. if layer.exllama_state == 0:
  103. layer.q_scale_max /= 256
  104. layer.q_invperm = layer.q_invperm.short()
  105. if not hasattr(layer, 'q_perm'):
  106. layer.q_perm = torch.argsort(layer.q_invperm).to(torch.short)
  107. if not hasattr(layer, 'q_group_map'):
  108. layer.q_group_map = make_group_map(layer.q_groups,
  109. layer.q_weight.shape[0])
  110. layer.q_matrix = ops.exl2_make_q_matrix(
  111. layer.q_weight,
  112. layer.q_perm,
  113. layer.q_invperm,
  114. layer.q_scale,
  115. layer.q_scale_max,
  116. layer.q_groups,
  117. layer.q_group_map,
  118. )
  119. layer.exllama_state = 1
  120. output = ops.exl2_gemm(reshaped_x, layer.q_matrix)
  121. if bias is not None:
  122. output.add_(bias)
  123. return output.reshape(out_shape)
  124. def apply_moe_weights(self, w1: Dict[str,
  125. torch.Tensor], w2: Dict[str,
  126. torch.Tensor],
  127. x: torch.Tensor, gating_output: torch.Tensor,
  128. topk: int, renormalize: bool) -> torch.Tensor:
  129. raise NotImplementedError