mirror of
https://github.com/marian-nmt/marian.git
synced 2024-08-15 17:20:28 +03:00
Update script exporting embeddings to support tied embeddings (#569)
This commit is contained in:
parent
22ad592a1d
commit
533604024b
@ -9,18 +9,22 @@ import numpy as np
|
||||
|
||||
|
||||
def main():
|
||||
desc = """Export word embedding from model"""
|
||||
desc = """Export word embeddings from model"""
|
||||
parser = argparse.ArgumentParser(
|
||||
formatter_class=argparse.RawDescriptionHelpFormatter, description=desc)
|
||||
parser.add_argument("-m", "--model", help="Model file", required=True)
|
||||
parser.add_argument(
|
||||
"-o", "--output-prefix", help="Output files prefix", required=True)
|
||||
parser.add_argument("-m", "--model", help="path to model.npz file", required=True)
|
||||
parser.add_argument("-o", "--output-prefix", help="prefix for output files", required=True)
|
||||
args = parser.parse_args()
|
||||
|
||||
print("Loading model")
|
||||
model = np.load(args.model)
|
||||
special = yaml.load(model["special:model.yml"][:-1].tobytes())
|
||||
|
||||
if special["tied-embeddings-all"] or special["tied-embeddings-src"]:
|
||||
all_emb = model["Wemb"]
|
||||
export_emb(args.output_prefix + ".all", all_emb)
|
||||
exit()
|
||||
|
||||
if special["type"] == "amun":
|
||||
enc_emb = model["Wemb"]
|
||||
dec_emb = model["Wemb_dec"]
|
||||
@ -28,16 +32,15 @@ def main():
|
||||
enc_emb = model["encoder_Wemb"]
|
||||
dec_emb = model["decoder_Wemb"]
|
||||
|
||||
with open(args.output_prefix + ".src", "w") as out:
|
||||
out.write("{0} {1}\n".format(*enc_emb.shape))
|
||||
for i in range(enc_emb.shape[0]):
|
||||
vec = " ".join("{0:.8f}".format(v) for v in enc_emb[i])
|
||||
out.write("{0} {1}\n".format(i, vec))
|
||||
export_emb(args.output_prefix + ".src", enc_emb)
|
||||
export_emb(args.output_prefix + ".trg", dec_emb)
|
||||
|
||||
with open(args.output_prefix + ".trg", "w") as out:
|
||||
out.write("{0} {1}\n".format(*dec_emb.shape))
|
||||
for i in range(dec_emb.shape[0]):
|
||||
vec = " ".join("{0:.8f}".format(v) for v in dec_emb[i])
|
||||
|
||||
def export_emb(filename, emb):
|
||||
with open(filename, "w") as out:
|
||||
out.write("{0} {1}\n".format(*emb.shape))
|
||||
for i in range(emb.shape[0]):
|
||||
vec = " ".join("{0:.8f}".format(v) for v in emb[i])
|
||||
out.write("{0} {1}\n".format(i, vec))
|
||||
|
||||
|
||||
|
Loading…
Reference in New Issue
Block a user