weights.py 21 KB
Newer Older
1
from dataclasses import dataclass, field
Nicolas Patry's avatar
Nicolas Patry committed
2
import os
3
from pathlib import Path
4
from typing import List, Dict, Optional, Set, Tuple, Union
5
from safetensors import safe_open, SafetensorError
6
import torch
7
from loguru import logger
8
9
from huggingface_hub import hf_hub_download
import json
10
from text_generation_server.utils.log import log_once
11
12
13


class Weights:
14
15
16
17
18
19
20
    def __init__(
        self,
        filenames: List[Path],
        device,
        dtype,
        process_group,
        aliases: Optional[Dict[str, List[str]]] = None,
OlivierDehaene's avatar
OlivierDehaene committed
21
        prefix: Optional[str] = None,
22
    ):
23
24
25
26
27
28
29
30
31
        routing = {}
        for filename in filenames:
            with safe_open(filename, framework="pytorch") as f:
                for k in f.keys():
                    if k in routing:
                        raise RuntimeError(
                            f"Key {k} was found in multiple files: {filename} and {routing[k]}"
                        )
                    routing[k] = filename
32
33
34
        if aliases is None:
            aliases = {}
        self.aliases = aliases
35
36
37
38
        self.routing = routing
        self.device = device
        self.dtype = dtype
        self.process_group = process_group
Nicolas Patry's avatar
Nicolas Patry committed
39
        self.prefix = prefix
40
41
42
43
44
45
46
47
48
        self._handles = {}

    def _get_handle(self, filename):
        if filename not in self._handles:
            f = safe_open(filename, framework="pytorch")
            self._handles[filename] = f

        return self._handles[filename]

49
    def get_filename(self, tensor_name: str) -> (str, str):
Nicolas Patry's avatar
Nicolas Patry committed
50
51
52
53
54
55
56
57
58
59
        names = [tensor_name]
        if self.prefix is not None:
            prefixed = f"{self.prefix}.{tensor_name}"
            names.append(prefixed)
        for name in names:
            filename = self.routing.get(name, None)
            if filename is not None:
                return str(filename), name

            aliases = self.aliases.get(name, [])
60
61
62
63
            for alias in aliases:
                filename = self.routing.get(alias, None)
                if filename is not None:
                    return str(filename), alias
Nicolas Patry's avatar
Nicolas Patry committed
64
        raise RuntimeError(f"weight {tensor_name} does not exist")
65
66

    def _get_slice(self, tensor_name: str):
67
        filename, tensor_name = self.get_filename(tensor_name)
68
69
70
71
72
73
74
        f = self._get_handle(filename)
        slice_ = f.get_slice(tensor_name)
        return slice_

    def get_shape(self, tensor_name: str):
        return self._get_slice(tensor_name).get_shape()

OlivierDehaene's avatar
OlivierDehaene committed
75
    def get_tensor(self, tensor_name: str, to_device=True):
76
        filename, tensor_name = self.get_filename(tensor_name)
77
78
        f = self._get_handle(filename)
        tensor = f.get_tensor(tensor_name)
79
        # Special case for gptq which shouldn't convert
80
81
82
        # u4 which are disguised as int32. Exl2 uses int16
        # as well.
        if tensor.dtype not in [torch.int16, torch.int32, torch.int64]:
83
            tensor = tensor.to(dtype=self.dtype)
xiaobin's avatar
xiaobin committed
84
85
        if to_device:
            tensor = tensor.to(device=self.device)
86
87
        return tensor

88
    def get_partial_sharded(self, tensor_name: str, dim: int):
89
        filename, tensor_name = self.get_filename(tensor_name)
xiaobin's avatar
xiaobin committed
90
91
        f = self._get_handle(filename)
        slice_ = f.get_slice(tensor_name)
92
93
94
95
        world_size = self.process_group.size()
        rank = self.process_group.rank()

        size = slice_.get_shape()[dim]
96
        block_size = (size + world_size - 1) // world_size
97
98
99
100
101
102
103
104
105
        start = rank * block_size
        stop = (rank + 1) * block_size

        if dim == 0:
            tensor = slice_[start:stop]
        elif dim == 1:
            tensor = slice_[:, start:stop]
        else:
            raise NotImplementedError("Let's make that generic when needed")
106
        # Special case for gptq which shouldn't convert
107
108
        # u4 which are disguised as int32. exl2 uses int16.
        if tensor.dtype not in (torch.int16, torch.int32):
109
            tensor = tensor.to(dtype=self.dtype)
110
111
        tensor = tensor.to(device=self.device)
        return tensor
112

113
114
115
116
117
118
119
120
121
122
123
    def get_sharded(self, tensor_name: str, dim: int):
        filename, tensor_name = self.get_filename(tensor_name)
        f = self._get_handle(filename)
        slice_ = f.get_slice(tensor_name)
        world_size = self.process_group.size()
        size = slice_.get_shape()[dim]
        assert (
            size % world_size == 0
        ), f"The choosen size {size} is not compatible with sharding on {world_size} shards"
        return self.get_partial_sharded(tensor_name, dim)

124
    def _get_qweight(self, name: str, blocks: int):
xiaobin's avatar
xiaobin committed
125
126
        slice_ = self._get_slice(name)
        total_size = slice_.get_shape()[1]
127
128
129
130
        assert (
            total_size % blocks == 0
        ), f"Prepacked quantized matrix is not divisible by {blocks}"
        single_size = total_size // blocks
xiaobin's avatar
xiaobin committed
131
132
133
        world_size = self.process_group.size()
        rank = self.process_group.rank()

OlivierDehaene's avatar
OlivierDehaene committed
134
135
        assert (
            single_size % world_size == 0
136
        ), f"Prepacked quantized matrix cannot be sharded across {world_size} shards"
xiaobin's avatar
xiaobin committed
137
138
139
        block_size = single_size // world_size
        start = rank * block_size
        stop = (rank + 1) * block_size
140
141
142
143
144
145
146
147

        weights = []
        for block in range(blocks):
            weights.append(
                slice_[:, start + block * single_size : stop + block * single_size]
            )

        weight = torch.cat(weights, dim=1)
xiaobin's avatar
xiaobin committed
148
149
150
151
        weight = weight.to(device=self.device)
        return weight

    def get_weights_col_packed_qkv(self, prefix: str, quantize: str):
Nicolas Patry's avatar
Nicolas Patry committed
152
153
154
155
156
157
        return self.get_weights_col_packed(prefix, quantize, 3)

    def get_weights_col_packed_gate_up(self, prefix: str, quantize: str):
        return self.get_weights_col_packed(prefix, quantize, 2)

    def get_weights_col_packed(self, prefix: str, quantize: str, blocks: int):
xiaobin's avatar
xiaobin committed
158
159
160
161
        """
        Highly specific when the underlying tensor is a simple cat of Q,K,V instead of being
        already alternating Q,K,V within the main tensor
        """
162
        if quantize in ["gptq", "awq"]:
Nicolas Patry's avatar
Nicolas Patry committed
163
164
            from text_generation_server.layers.gptq import GPTQWeight

xiaobin's avatar
xiaobin committed
165
            try:
166
                qweight = self._get_qweight(f"{prefix}.qweight", blocks)
xiaobin's avatar
xiaobin committed
167
168
            except RuntimeError:
                raise RuntimeError(
169
                    f"Cannot load `{quantize}` weight, make sure the model is already quantized."
xiaobin's avatar
xiaobin committed
170
171
                )

Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
172
173
            bits, groupsize, _, quant_method = self._get_gptq_params()

174
175
            qzeros = self._get_qweight(f"{prefix}.qzeros", blocks)
            scales = self._get_qweight(f"{prefix}.scales", blocks)
xiaobin's avatar
xiaobin committed
176
            scales = scales.to(dtype=self.dtype)
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
177
178

            if quantize == "gptq" and quant_method == "gptq":
179
                g_idx = self.get_tensor(f"{prefix}.g_idx")
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
180
181
182
183
            elif quantize == "gptq" and quant_method == "awq":
                log_once(
                    logger.info, "Converting AWQ model to Exllama/GPTQ packing format."
                )
Nicolas Patry's avatar
Nicolas Patry committed
184
                from text_generation_server.layers.awq.conversion_utils import (
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
185
186
187
188
189
190
191
192
                    fast_awq_to_gptq,
                )

                qweight, qzeros = fast_awq_to_gptq(qweight, qzeros)
                g_idx = (
                    torch.arange(qweight.shape[0] * (32 // bits), device=qweight.device)
                    // groupsize
                ).to(dtype=torch.int32)
193
194
            else:
                g_idx = None
xiaobin's avatar
xiaobin committed
195

196
197
198
199
200
201
202
203
204
            weight = GPTQWeight(
                qweight=qweight,
                qzeros=qzeros,
                scales=scales,
                g_idx=g_idx,
                bits=bits,
                groupsize=groupsize,
                use_exllama=False,
            )
xiaobin's avatar
xiaobin committed
205
        else:
OlivierDehaene's avatar
OlivierDehaene committed
206
            slice_ = self._get_slice(f"{prefix}.weight")
xiaobin's avatar
xiaobin committed
207
            total_size = slice_.get_shape()[0]
Nicolas Patry's avatar
Nicolas Patry committed
208
209
            assert total_size % blocks == 0, f"Prepacked is not divisible by {blocks}"
            single_size = total_size // blocks
xiaobin's avatar
xiaobin committed
210
211
212
            world_size = self.process_group.size()
            rank = self.process_group.rank()

OlivierDehaene's avatar
OlivierDehaene committed
213
214
215
            assert (
                single_size % world_size == 0
            ), f"Prepacked qkv cannot be sharded across {world_size} shards"
xiaobin's avatar
xiaobin committed
216
217
218
            block_size = single_size // world_size
            start = rank * block_size
            stop = (rank + 1) * block_size
Nicolas Patry's avatar
Nicolas Patry committed
219
220
221
222
223
            tensors = []
            for i in range(blocks):
                tensor = slice_[start + i * single_size : stop + i * single_size]
                tensors.append(tensor)
            weight = torch.cat(tensors, dim=0)
xiaobin's avatar
xiaobin committed
224
225
226
227
            weight = weight.to(device=self.device)
            weight = weight.to(dtype=self.dtype)
        return weight

228
229
    def get_weights_col(self, prefix: str, quantize: str):
        if quantize == "exl2":
Nicolas Patry's avatar
Nicolas Patry committed
230
231
            from text_generation_server.layers.exl2 import Exl2Weight

232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
            try:
                q_weight = self.get_tensor(f"{prefix}.q_weight")
            except RuntimeError:
                raise RuntimeError(
                    f"Cannot load `exl2`-quantized weight, make sure the model is already quantized."
                )

            q_scale = self.get_tensor(f"{prefix}.q_scale")
            q_invperm = self.get_tensor(f"{prefix}.q_invperm")
            q_scale_max = self.get_tensor(f"{prefix}.q_scale_max")
            q_groups = self.get_tensor(f"{prefix}.q_groups")

            return Exl2Weight(
                q_weight=q_weight,
                q_scale=q_scale,
                q_invperm=q_invperm,
                q_scale_max=q_scale_max,
                q_groups=q_groups,
            )

        return self.get_multi_weights_col([prefix], quantize, 0)

254
    def get_multi_weights_col(self, prefixes: List[str], quantize: str, dim: int):
255
256
257
        if quantize == "exl2":
            raise ValueError("get_multi_weights_col is not supported for exl2")
        elif quantize in ["gptq", "awq"]:
Nicolas Patry's avatar
Nicolas Patry committed
258
259
            from text_generation_server.layers.gptq import GPTQWeight

260
            try:
261
262
263
                qweight = torch.cat(
                    [self.get_sharded(f"{p}.qweight", dim=1) for p in prefixes], dim=1
                )
264
            except RuntimeError:
265
                raise RuntimeError(
266
                    f"Cannot load `{quantize}` weight, make sure the model is already quantized"
267
268
269
270
271
272
273
274
                )

            qzeros = torch.cat(
                [self.get_sharded(f"{p}.qzeros", dim=1) for p in prefixes], dim=1
            )
            scales = torch.cat(
                [self.get_sharded(f"{p}.scales", dim=1) for p in prefixes], dim=1
            )
275

Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
276
277
            bits, groupsize, desc_act, quant_method = self._get_gptq_params()

Nicolas Patry's avatar
Nicolas Patry committed
278
            from text_generation_server.layers.gptq import HAS_EXLLAMA
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
279
280
281
282
283
284

            use_exllama = (
                bits == 4 and HAS_EXLLAMA and quantize == "gptq" and not desc_act
            )

            if quantize == "gptq" and quant_method == "gptq":
285
286
287
288
                w = [self.get_tensor(f"{p}.g_idx") for p in prefixes]
                for w2 in w[1:]:
                    torch.testing.assert_close(w2, w[0])
                g_idx = w[0]
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
289
290
291
292
            elif quantize == "gptq" and quant_method == "awq":
                log_once(
                    logger.info, "Converting AWQ model to Exllama/GPTQ packing format."
                )
Nicolas Patry's avatar
Nicolas Patry committed
293
                from text_generation_server.layers.awq.conversion_utils import (
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
294
295
296
297
298
299
300
301
302
303
304
305
306
                    fast_awq_to_gptq,
                )

                qweight, qzeros = fast_awq_to_gptq(qweight, qzeros)
                if use_exllama:
                    g_idx = None
                else:
                    g_idx = (
                        torch.arange(
                            qweight.shape[0] * (32 // bits), device=qweight.device
                        )
                        // groupsize
                    ).to(dtype=torch.int32)
307
308
            else:
                g_idx = None
309

310
311
312
313
314
315
316
317
318
            weight = GPTQWeight(
                qweight=qweight,
                qzeros=qzeros,
                scales=scales,
                g_idx=g_idx,
                bits=bits,
                groupsize=groupsize,
                use_exllama=use_exllama,
            )
319
320
321
322
        else:
            w = [self.get_sharded(f"{p}.weight", dim=0) for p in prefixes]
            weight = torch.cat(w, dim=dim)
        return weight
OlivierDehaene's avatar
OlivierDehaene committed
323

xiaobin's avatar
xiaobin committed
324
325
326
327
328
329
330
331
332
333
334
335
336
337
    def get_tensor_shard(self, var, dim):
        world_size = self.process_group.size()
        rank = self.process_group.rank()
        block_size = var.size()[dim] // world_size
        start = rank * block_size
        stop = (rank + 1) * block_size
        if dim == 0:
            tensor = var[start:stop]
        elif dim == 1:
            tensor = var[:, start:stop]
        else:
            raise NotImplementedError("Let's make that generic when needed")
        tensor = tensor.to(dtype=self.dtype)
        tensor = tensor.to(device=self.device)
OlivierDehaene's avatar
OlivierDehaene committed
338
        return tensor
339
340

    def get_multi_weights_row(self, prefix: str, quantize: str):
341
        if quantize == "exl2":
Nicolas Patry's avatar
Nicolas Patry committed
342
343
            from text_generation_server.layers.exl2 import Exl2Weight

344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
            try:
                q_weight = self.get_tensor(f"{prefix}.q_weight")
            except RuntimeError:
                raise RuntimeError(
                    f"Cannot load `exl2`-quantized weight, make sure the model is already quantized."
                )

            q_scale = self.get_tensor(f"{prefix}.q_scale")
            q_invperm = self.get_tensor(f"{prefix}.q_invperm")
            q_scale_max = self.get_tensor(f"{prefix}.q_scale_max")
            q_groups = self.get_tensor(f"{prefix}.q_groups")

            return Exl2Weight(
                q_weight=q_weight,
                q_scale=q_scale,
                q_invperm=q_invperm,
                q_scale_max=q_scale_max,
                q_groups=q_groups,
            )

        elif quantize == "gptq":
365
            use_exllama = True
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
366
            bits, groupsize, desc_act, quant_method = self._get_gptq_params()
367
368
369
370

            if bits != 4:
                use_exllama = False

371
372
373
374
            if desc_act:
                log_once(logger.warning, "Disabling exllama because desc_act=True")
                use_exllama = False

Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
375
376
377
378
379
380
381
382
383
384
385
386
            try:
                qweight = self.get_sharded(f"{prefix}.qweight", dim=0)
            except RuntimeError:
                raise RuntimeError(
                    "Cannot load `gptq` weight, make sure the model is already quantized, or quantize it with `text-generation-server quantize ORIGINAL_MODEL_ID NEW_MODEL_ID`"
                )

            if quant_method == "gptq":
                g_idx = self.get_sharded(f"{prefix}.g_idx", dim=0)
            elif quant_method == "awq":
                g_idx = None

387
388
            if self.process_group.size() > 1:
                if g_idx is not None:
389
390
391
392
393
394
395
396
397
398
                    if (
                        not torch.equal(
                            g_idx.cpu(),
                            torch.tensor(
                                [i // groupsize for i in range(g_idx.shape[0])],
                                dtype=torch.int32,
                            ),
                        )
                        and not (g_idx == 0).all()
                    ):
399
400
401
402
                        # Exllama implementation does not support row tensor parallelism with act-order, as
                        # it would require to reorder input activations that are split unto several GPUs
                        use_exllama = False

Nicolas Patry's avatar
Nicolas Patry committed
403
404
405
406
407
            from text_generation_server.layers.gptq import (
                HAS_EXLLAMA,
                CAN_EXLLAMA,
                GPTQWeight,
            )
408

409
            if use_exllama:
410
411
                if not HAS_EXLLAMA:
                    if CAN_EXLLAMA:
412
413
                        log_once(
                            logger.warning,
OlivierDehaene's avatar
v1.3.4  
OlivierDehaene committed
414
                            "Exllama GPTQ cuda kernels (which are faster) could have been used, but are not currently installed, try using BUILD_EXTENSIONS=True",
415
                        )
416
417
                    use_exllama = False
                else:
OlivierDehaene's avatar
v1.3.4  
OlivierDehaene committed
418
                    log_once(logger.info, f"Using exllama kernels v{HAS_EXLLAMA}")
419

420
            if use_exllama and groupsize != -1:
Nicolas Patry's avatar
Nicolas Patry committed
421
422
                qzeros = self.get_sharded(f"{prefix}.qzeros", dim=0)
                scales = self.get_sharded(f"{prefix}.scales", dim=0)
423
424
425
            else:
                qzeros = self.get_tensor(f"{prefix}.qzeros")
                scales = self.get_tensor(f"{prefix}.scales")
426

Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
427
            if use_exllama and g_idx is not None:
428
                g_idx = g_idx - g_idx[0]
429

Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
430
431
432
433
            if quant_method == "awq":
                log_once(
                    logger.info, "Converting AWQ model to Exllama/GPTQ packing format."
                )
Nicolas Patry's avatar
Nicolas Patry committed
434
                from text_generation_server.layers.awq.conversion_utils import (
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
435
436
437
438
439
440
441
442
443
444
445
446
447
448
                    fast_awq_to_gptq,
                )

                qweight, qzeros = fast_awq_to_gptq(qweight, qzeros)
                if use_exllama:
                    g_idx = None
                else:
                    g_idx = (
                        torch.arange(
                            qweight.shape[0] * (32 // bits), device=qweight.device
                        )
                        // groupsize
                    ).to(dtype=torch.int32)

449
450
451
452
453
454
455
456
457
            weight = GPTQWeight(
                qweight=qweight,
                qzeros=qzeros,
                scales=scales,
                g_idx=g_idx,
                bits=bits,
                groupsize=groupsize,
                use_exllama=use_exllama,
            )
458
        elif quantize == "awq":
Nicolas Patry's avatar
Nicolas Patry committed
459
460
            from text_generation_server.layers.gptq import GPTQWeight

Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
461
            bits, groupsize, _, _ = self._get_gptq_params()
462
463
464
465
466
467
468
469
470
471
472
473

            try:
                qweight = self.get_sharded(f"{prefix}.qweight", dim=0)
            except RuntimeError:
                raise RuntimeError(
                    "Cannot load `awq` weight, make sure the model is already quantized"
                )

            qzeros = self.get_sharded(f"{prefix}.qzeros", dim=0)
            scales = self.get_sharded(f"{prefix}.scales", dim=0)
            g_idx = None
            use_exllama = False
OlivierDehaene's avatar
OlivierDehaene committed
474

475
476
477
478
479
480
481
482
483
            weight = GPTQWeight(
                qweight=qweight,
                qzeros=qzeros,
                scales=scales,
                g_idx=g_idx,
                bits=bits,
                groupsize=groupsize,
                use_exllama=use_exllama,
            )
484
485
486
        else:
            weight = self.get_sharded(f"{prefix}.weight", dim=1)
        return weight
487

Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
488
    def _get_gptq_params(self) -> Tuple[int, int, int, str]:
489
490
491
        try:
            bits = self.get_tensor("gptq_bits").item()
            groupsize = self.get_tensor("gptq_groupsize").item()
492
            desc_act = False
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
493
            quant_method = "gptq"
494
495
        except (SafetensorError, RuntimeError) as e:
            try:
496
497
                bits = self.gptq_bits
                groupsize = self.gptq_groupsize
498
                desc_act = getattr(self, "gptq_desc_act", False)
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
499
                quant_method = getattr(self, "quant_method", "gptq")
500
501
502
            except Exception:
                raise e

Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
503
        return bits, groupsize, desc_act, quant_method
504

OlivierDehaene's avatar
OlivierDehaene committed
505
    def _set_gptq_params(self, model_id, revision):
506
        filename = "config.json"
507
        try:
508
            if os.path.exists(os.path.join(model_id, filename)):
Nicolas Patry's avatar
Nicolas Patry committed
509
510
                filename = os.path.join(model_id, filename)
            else:
OlivierDehaene's avatar
OlivierDehaene committed
511
512
513
                filename = hf_hub_download(
                    model_id, filename=filename, revision=revision
                )
514
515
            with open(filename, "r") as f:
                data = json.load(f)
516
517
            self.gptq_bits = data["quantization_config"]["bits"]
            self.gptq_groupsize = data["quantization_config"]["group_size"]
518
            # Order is important here, desc_act is missing on some real models
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
519
            self.quant_method = data["quantization_config"]["quant_method"]
520
            self.gptq_desc_act = data["quantization_config"]["desc_act"]
521
        except Exception:
522
523
524
525
526
            filename = "quantize_config.json"
            try:
                if os.path.exists(os.path.join(model_id, filename)):
                    filename = os.path.join(model_id, filename)
                else:
OlivierDehaene's avatar
OlivierDehaene committed
527
528
529
                    filename = hf_hub_download(
                        model_id, filename=filename, revision=revision
                    )
530
531
532
533
                with open(filename, "r") as f:
                    data = json.load(f)
                self.gptq_bits = data["bits"]
                self.gptq_groupsize = data["group_size"]
534
                self.gptq_desc_act = data["desc_act"]
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
535
536
                if "version" in data and data["version"] == "GEMM":
                    self.quant_method = "awq"
537
            except Exception:
538
539
540
541
542
                filename = "quant_config.json"
                try:
                    if os.path.exists(os.path.join(model_id, filename)):
                        filename = os.path.join(model_id, filename)
                    else:
OlivierDehaene's avatar
OlivierDehaene committed
543
544
545
                        filename = hf_hub_download(
                            model_id, filename=filename, revision=revision
                        )
546
547
548
549
                    with open(filename, "r") as f:
                        data = json.load(f)
                    self.gptq_bits = data["w_bit"]
                    self.gptq_groupsize = data["q_group_size"]
550
                    self.gptq_desc_act = data["desc_act"]
Ilyas Moutawwakil's avatar
Ilyas Moutawwakil committed
551
552
                    if "version" in data and data["version"] == "GEMM":
                        self.quant_method = "awq"
553
554
                except Exception:
                    pass